Skip to content

Commit c138173

Browse files
committed
fix mypy and flake
1 parent e69ff3b commit c138173

11 files changed

+68
-46
lines changed

autoPyTorch/api/base_task.py

+1-1
Original file line numberDiff line numberDiff line change
@@ -1908,7 +1908,7 @@ def _init_ensemble_builder(
19081908
# builder in the provide dask client
19091909
required_dataset_properties = {'task_type': self.task_type,
19101910
'output_type': self.dataset.output_type}
1911-
1911+
19121912
proc_ensemble = EnsembleBuilderManager(
19131913
start_time=time.time(),
19141914
time_left_for_ensembles=time_left_for_ensembles,

autoPyTorch/api/time_series_forecasting.py

+6
Original file line numberDiff line numberDiff line change
@@ -526,6 +526,9 @@ def predict(
526526
predicted value, it needs to be with shape (B, H, N),
527527
B is the number of series, H is forecasting horizon (n_prediction_steps), N is the number of targets
528528
"""
529+
if self.dataset is None:
530+
raise AttributeError(f"Expected dataset to be initialised when predicting in {self.__class__.__name__}")
531+
529532
if X_test is None or not isinstance(X_test[0], TimeSeriesSequence):
530533
assert past_targets is not None
531534
# Validate and construct TimeSeriesSequence
@@ -566,6 +569,9 @@ def update_sliding_window_size(self, n_prediction_steps: int) -> None:
566569
forecast horizon. Sometimes we could also make our base sliding window size based on the
567570
forecast horizon
568571
"""
572+
if self.dataset is None:
573+
raise AttributeError(f"Expected dataset to be initialised when updating sliding window"
574+
f" in {self.__class__.__name__}")
569575
base_window_size = int(np.ceil(self.dataset.base_window_size))
570576
# we don't want base window size to large, which might cause a too long computation time, in which case
571577
# we will use n_prediction_step instead (which is normally smaller than base_window_size)

autoPyTorch/data/base_feature_validator.py

+1-1
Original file line numberDiff line numberDiff line change
@@ -49,7 +49,7 @@ def __init__(
4949
self.categories: List[List[int]] = []
5050
self.categorical_columns: List[int] = []
5151
self.numerical_columns: List[int] = []
52-
self.encode_columns: List[int] = []
52+
self.encode_columns: List[str] = []
5353

5454
self.all_nan_columns: Optional[Set[Union[int, str]]] = None
5555

autoPyTorch/data/tabular_feature_validator.py

+1-2
Original file line numberDiff line numberDiff line change
@@ -283,7 +283,6 @@ def transform(
283283
X = self.numpy_to_pandas(X)
284284

285285
if ispandas(X) and not issparse(X):
286-
X = cast(pd.DataFrame, X)
287286

288287
if self.all_nan_columns is None:
289288
raise ValueError('_fit must be called before calling transform')
@@ -491,7 +490,7 @@ def _get_columns_to_encode(
491490
# Also, register the feature types for the estimator
492491
feat_types = []
493492

494-
# Make sure each column is a valid type
493+
# Make sure each column is a valid type
495494
for i, column in enumerate(X.columns):
496495
if self.all_nan_columns is not None and column in self.all_nan_columns:
497496
continue

autoPyTorch/optimizer/smbo.py

+2-1
Original file line numberDiff line numberDiff line change
@@ -276,7 +276,8 @@ def __init__(self,
276276
initial_configurations = []
277277

278278
if STRING_TO_TASK_TYPES.get(self.task_type, -1) == TIMESERIES_FORECASTING:
279-
# TODO: update search space (to remove reg cocktails) for forecasting tasks so that we can use the portfolio (or build the portfolio again)
279+
# TODO: update search space (to remove reg cocktails) for forecasting tasks so
280+
# that we can use the portfolio (or build the portfolio again)
280281
# initial_configurations = self.get_init_configs_for_forecasting(config_space, kwargs)
281282
# proxy-validation sets
282283
self.min_num_test_instances: Optional[int] = kwargs.get('min_num_test_instances', # type:ignore[assignment]

autoPyTorch/pipeline/base_pipeline.py

+6-2
Original file line numberDiff line numberDiff line change
@@ -1,7 +1,7 @@
1-
from copy import copy
21
import warnings
32
from abc import ABCMeta
43
from collections import Counter
4+
from copy import copy
55
from typing import Any, Dict, List, Optional, Tuple, Union
66

77
from ConfigSpace import Configuration
@@ -297,7 +297,7 @@ def _get_hyperparameter_search_space(self,
297297
"""
298298
raise NotImplementedError()
299299

300-
def _add_forbidden_conditions(self, cs):
300+
def _add_forbidden_conditions(self, cs: ConfigurationSpace) -> ConfigurationSpace:
301301
"""
302302
Add forbidden conditions to ensure valid configurations.
303303
Currently, Learned Entity Embedding is only valid when encoder is one hot encoder
@@ -308,6 +308,10 @@ def _add_forbidden_conditions(self, cs):
308308
cs (ConfigurationSpace):
309309
Configuration space to which forbidden conditions are added.
310310
311+
Returns:
312+
ConfigurationSpace:
313+
with forbidden conditions added to the search space
314+
311315
"""
312316

313317
# Learned Entity Embedding is only valid when encoder is one hot encoder

autoPyTorch/pipeline/components/setup/network/forecasting_architecture.py

+47-13
Original file line numberDiff line numberDiff line change
@@ -576,7 +576,14 @@ def forward(self,
576576
) -> ALL_NET_OUTPUT:
577577

578578
if isinstance(past_targets, dict):
579-
past_targets, past_features, future_features, past_observed_targets = self._unwrap_past_targets(past_targets)
579+
(
580+
past_targets,
581+
past_features,
582+
future_features,
583+
past_observed_targets,
584+
future_targets,
585+
decoder_observed_values
586+
) = self._unwrap_past_targets(past_targets)
580587

581588
x_past, x_future, x_static, loc, scale, static_context_initial_hidden, _ = self.pre_processing(
582589
past_targets=past_targets,
@@ -610,13 +617,12 @@ def forward(self,
610617
def _unwrap_past_targets(
611618
self,
612619
past_targets: dict
613-
) -> Tuple[
614-
torch.Tensor,
615-
Optional[torch.Tensor],
616-
Optional[torch.Tensor],
617-
Optional[torch.Tensor],
618-
Optional[torch.BoolTensor],
619-
Optional[torch.Tensor]]:
620+
) -> Tuple[torch.Tensor,
621+
Optional[torch.Tensor],
622+
Optional[torch.Tensor],
623+
Optional[torch.Tensor],
624+
Optional[torch.BoolTensor],
625+
Optional[torch.Tensor]]:
620626
"""
621627
Time series forecasting network requires multiple inputs for the forward pass which is different to how pytorch
622628
networks usually work. SWA's update_bn in line #452 of trainer choice, does not unwrap the dictionary of the
@@ -637,7 +643,14 @@ def _unwrap_past_targets(
637643
future_features = past_targets_copy.pop('future_features', None)
638644
past_observed_targets = past_targets_copy.pop('past_observed_targets', None)
639645
decoder_observed_values = past_targets_copy.pop('decoder_observed_values', None)
640-
return past_targets,past_features,future_features,past_observed_targets
646+
return (
647+
past_targets,
648+
past_features,
649+
future_features,
650+
past_observed_targets,
651+
future_targets,
652+
decoder_observed_values
653+
)
641654

642655
def pred_from_net_output(self, net_output: ALL_NET_OUTPUT) -> torch.Tensor:
643656
if self.output_type == 'regression':
@@ -730,9 +743,16 @@ def forward(self,
730743
future_features: Optional[torch.Tensor] = None,
731744
past_observed_targets: Optional[torch.BoolTensor] = None,
732745
decoder_observed_values: Optional[torch.Tensor] = None, ) -> ALL_NET_OUTPUT:
733-
746+
734747
if isinstance(past_targets, dict):
735-
past_targets, past_features, future_features, past_observed_targets = self._unwrap_past_targets(past_targets)
748+
(
749+
past_targets,
750+
past_features,
751+
future_features,
752+
past_observed_targets,
753+
future_targets,
754+
decoder_observed_values
755+
) = self._unwrap_past_targets(past_targets)
736756

737757
x_past, _, x_static, loc, scale, static_context_initial_hidden, past_targets = self.pre_processing(
738758
past_targets=past_targets,
@@ -1025,7 +1045,14 @@ def forward(self,
10251045
decoder_observed_values: Optional[torch.Tensor] = None, ) -> ALL_NET_OUTPUT:
10261046

10271047
if isinstance(past_targets, dict):
1028-
past_targets, past_features, future_features, past_observed_targets = self._unwrap_past_targets(past_targets)
1048+
(
1049+
past_targets,
1050+
past_features,
1051+
future_features,
1052+
past_observed_targets,
1053+
future_targets,
1054+
decoder_observed_values
1055+
) = self._unwrap_past_targets(past_targets)
10291056

10301057
encode_length = min(self.window_size, past_targets.shape[1])
10311058

@@ -1295,7 +1322,14 @@ def forward(self, # type: ignore[override]
12951322
Tuple[torch.Tensor, torch.Tensor]]:
12961323

12971324
if isinstance(past_targets, dict):
1298-
past_targets, past_features, future_features, past_observed_targets = self._unwrap_past_targets(past_targets)
1325+
(
1326+
past_targets,
1327+
past_features,
1328+
future_features,
1329+
past_observed_targets,
1330+
future_targets,
1331+
decoder_observed_values
1332+
) = self._unwrap_past_targets(past_targets)
12991333

13001334
# Unlike other networks, NBEATS network is required to predict both past and future targets.
13011335
# Thereby, we return two tensors for backcast and forecast

autoPyTorch/pipeline/tabular_classification.py

-1
Original file line numberDiff line numberDiff line change
@@ -1,4 +1,3 @@
1-
import copy
21
import warnings
32
from typing import Any, Dict, List, Optional, Tuple, Union
43

autoPyTorch/pipeline/tabular_regression.py

-1
Original file line numberDiff line numberDiff line change
@@ -1,4 +1,3 @@
1-
import copy
21
import warnings
32
from typing import Any, Dict, List, Optional, Tuple, Union
43

test/test_data/test_feature_validator.py

-23
Original file line numberDiff line numberDiff line change
@@ -645,7 +645,6 @@ def test_feature_validator_get_columns_to_encode_error_feat_type(input_data_feat
645645
with pytest.raises(ValueError, match=r"Expected type of features to be in .*"):
646646
validator._validate_feat_types(X)
647647

648-
649648
# Null columns in the train split but not necessarily in the test split
650649
train_features = {
651650
'A': [np.NaN, np.NaN, np.NaN],
@@ -706,25 +705,3 @@ def test_feature_validator_get_columns_to_encode_error_feat_type(input_data_feat
706705
null_columns.append(column)
707706

708707
assert null_columns == [1]
709-
710-
def test_comparator():
711-
numerical = 'numerical'
712-
categorical = 'categorical'
713-
714-
validator = TabularFeatureValidator
715-
716-
feat_type = [numerical, categorical] * 10
717-
ans = [categorical] * 10 + [numerical] * 10
718-
feat_type = sorted(
719-
feat_type,
720-
key=functools.cmp_to_key(validator._comparator)
721-
)
722-
assert ans == feat_type
723-
724-
feat_type = [numerical] * 10 + [categorical] * 10
725-
ans = [categorical] * 10 + [numerical] * 10
726-
feat_type = sorted(
727-
feat_type,
728-
key=functools.cmp_to_key(validator._comparator)
729-
)
730-
assert ans == feat_type

test/test_pipeline/test_tabular_classification.py

+4-1
Original file line numberDiff line numberDiff line change
@@ -35,7 +35,10 @@
3535

3636
@pytest.fixture
3737
def exclude():
38-
return {'feature_preprocessor': ['SelectRatesClassification', 'SelectPercentileClassification'], 'network_embedding': ['LearnedEntityEmbedding']}
38+
return {
39+
'feature_preprocessor': ['SelectRatesClassification', 'SelectPercentileClassification'],
40+
'network_embedding': ['LearnedEntityEmbedding']
41+
}
3942

4043

4144
@pytest.mark.parametrize("fit_dictionary_tabular", ['classification_categorical_only',

0 commit comments

Comments
 (0)