Skip to content

Commit b271e20

Browse files
XavierSATTLERglemaitre
authored andcommitted
DEP remove positive parameter for lars solver (scikit-learn#13863)
1 parent 7ea7284 commit b271e20

File tree

4 files changed

+130
-85
lines changed

4 files changed

+130
-85
lines changed

sklearn/decomposition/dict_learning.py

Lines changed: 17 additions & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -21,6 +21,14 @@
2121
from ..linear_model import Lasso, orthogonal_mp_gram, LassoLars, Lars
2222

2323

24+
def _check_positive_coding(method, positive):
25+
if positive and method in ["omp", "lars"]:
26+
raise ValueError(
27+
"Positive constraint not supported for '{}' "
28+
"coding method.".format(method)
29+
)
30+
31+
2432
def _sparse_encode(X, dictionary, gram, cov=None, algorithm='lasso_lars',
2533
regularization=None, copy_cov=True,
2634
init=None, max_iter=1000, check_input=True, verbose=0,
@@ -107,6 +115,8 @@ def _sparse_encode(X, dictionary, gram, cov=None, algorithm='lasso_lars',
107115
copy_cov = False
108116
cov = np.dot(dictionary, X.T)
109117

118+
_check_positive_coding(algorithm, positive)
119+
110120
if algorithm == 'lasso_lars':
111121
alpha = float(regularization) / n_features # account for scaling
112122
try:
@@ -147,7 +157,7 @@ def _sparse_encode(X, dictionary, gram, cov=None, algorithm='lasso_lars',
147157
# corrects the verbosity level.
148158
lars = Lars(fit_intercept=False, verbose=verbose, normalize=False,
149159
precompute=gram, n_nonzero_coefs=int(regularization),
150-
fit_path=False, positive=positive)
160+
fit_path=False)
151161
lars.fit(dictionary.T, X.T, Xy=cov)
152162
new_code = lars.coef_
153163
finally:
@@ -160,11 +170,6 @@ def _sparse_encode(X, dictionary, gram, cov=None, algorithm='lasso_lars',
160170
np.clip(new_code, 0, None, out=new_code)
161171

162172
elif algorithm == 'omp':
163-
# TODO: Should verbose argument be passed to this?
164-
if positive:
165-
raise ValueError(
166-
"Positive constraint not supported for \"omp\" coding method."
167-
)
168173
new_code = orthogonal_mp_gram(
169174
Gram=gram, Xy=cov, n_nonzero_coefs=int(regularization),
170175
tol=None, norms_squared=row_norms(X, squared=True),
@@ -519,6 +524,9 @@ def dict_learning(X, n_components, alpha, max_iter=100, tol=1e-8,
519524
if method not in ('lars', 'cd'):
520525
raise ValueError('Coding method %r not supported as a fit algorithm.'
521526
% method)
527+
528+
_check_positive_coding(method, positive_code)
529+
522530
method = 'lasso_' + method
523531

524532
t0 = time.time()
@@ -729,6 +737,9 @@ def dict_learning_online(X, n_components=2, alpha=1, n_iter=100,
729737

730738
if method not in ('lars', 'cd'):
731739
raise ValueError('Coding method not supported as a fit algorithm.')
740+
741+
_check_positive_coding(method, positive_code)
742+
732743
method = 'lasso_' + method
733744

734745
t0 = time.time()

sklearn/decomposition/tests/test_dict_learning.py

Lines changed: 99 additions & 46 deletions
Original file line numberDiff line numberDiff line change
@@ -18,6 +18,7 @@
1818
from sklearn.decomposition import DictionaryLearning
1919
from sklearn.decomposition import MiniBatchDictionaryLearning
2020
from sklearn.decomposition import SparseCoder
21+
from sklearn.decomposition import dict_learning
2122
from sklearn.decomposition import dict_learning_online
2223
from sklearn.decomposition import sparse_encode
2324

@@ -56,29 +57,30 @@ def test_dict_learning_overcomplete():
5657
assert dico.components_.shape == (n_components, n_features)
5758

5859

59-
# positive lars deprecated 0.22
60-
@pytest.mark.filterwarnings('ignore::DeprecationWarning')
60+
def test_dict_learning_lars_positive_parameter():
61+
n_components = 5
62+
alpha = 1
63+
err_msg = "Positive constraint not supported for 'lars' coding method."
64+
with pytest.raises(ValueError, match=err_msg):
65+
dict_learning(X, n_components, alpha, positive_code=True)
66+
67+
6168
@pytest.mark.parametrize("transform_algorithm", [
6269
"lasso_lars",
6370
"lasso_cd",
64-
"lars",
6571
"threshold",
6672
])
67-
@pytest.mark.parametrize("positive_code", [
68-
False,
69-
True,
70-
])
71-
@pytest.mark.parametrize("positive_dict", [
72-
False,
73-
True,
74-
])
73+
@pytest.mark.parametrize("positive_code", [False, True])
74+
@pytest.mark.parametrize("positive_dict", [False, True])
7575
def test_dict_learning_positivity(transform_algorithm,
7676
positive_code,
7777
positive_dict):
7878
n_components = 5
7979
dico = DictionaryLearning(
8080
n_components, transform_algorithm=transform_algorithm, random_state=0,
81-
positive_code=positive_code, positive_dict=positive_dict).fit(X)
81+
positive_code=positive_code, positive_dict=positive_dict,
82+
fit_algorithm="cd").fit(X)
83+
8284
code = dico.transform(X)
8385
if positive_dict:
8486
assert (dico.components_ >= 0).all()
@@ -90,6 +92,31 @@ def test_dict_learning_positivity(transform_algorithm,
9092
assert (code < 0).any()
9193

9294

95+
@pytest.mark.parametrize("positive_dict", [False, True])
96+
def test_dict_learning_lars_dict_positivity(positive_dict):
97+
n_components = 5
98+
dico = DictionaryLearning(
99+
n_components, transform_algorithm="lars", random_state=0,
100+
positive_dict=positive_dict, fit_algorithm="cd").fit(X)
101+
102+
if positive_dict:
103+
assert (dico.components_ >= 0).all()
104+
else:
105+
assert (dico.components_ < 0).any()
106+
107+
108+
def test_dict_learning_lars_code_positivity():
109+
n_components = 5
110+
dico = DictionaryLearning(
111+
n_components, transform_algorithm="lars", random_state=0,
112+
positive_code=True, fit_algorithm="cd").fit(X)
113+
114+
err_msg = "Positive constraint not supported for '{}' coding method."
115+
err_msg = err_msg.format("lars")
116+
with pytest.raises(ValueError, match=err_msg):
117+
dico.transform(X)
118+
119+
93120
def test_dict_learning_reconstruction():
94121
n_components = 12
95122
dico = DictionaryLearning(n_components, transform_algorithm='omp',
@@ -170,31 +197,29 @@ def test_dict_learning_online_shapes():
170197
assert_equal(np.dot(code, dictionary).shape, X.shape)
171198

172199

173-
# positive lars deprecated 0.22
174-
@pytest.mark.filterwarnings('ignore::DeprecationWarning')
200+
def test_dict_learning_online_lars_positive_parameter():
201+
alpha = 1
202+
err_msg = "Positive constraint not supported for 'lars' coding method."
203+
with pytest.raises(ValueError, match=err_msg):
204+
dict_learning_online(X, alpha, positive_code=True)
205+
206+
175207
@pytest.mark.parametrize("transform_algorithm", [
176208
"lasso_lars",
177209
"lasso_cd",
178-
"lars",
179210
"threshold",
180211
])
181-
@pytest.mark.parametrize("positive_code", [
182-
False,
183-
True,
184-
])
185-
@pytest.mark.parametrize("positive_dict", [
186-
False,
187-
True,
188-
])
189-
def test_dict_learning_online_positivity(transform_algorithm,
190-
positive_code,
191-
positive_dict):
192-
rng = np.random.RandomState(0)
212+
@pytest.mark.parametrize("positive_code", [False, True])
213+
@pytest.mark.parametrize("positive_dict", [False, True])
214+
def test_minibatch_dictionary_learning_positivity(transform_algorithm,
215+
positive_code,
216+
positive_dict):
193217
n_components = 8
194-
195218
dico = MiniBatchDictionaryLearning(
196219
n_components, transform_algorithm=transform_algorithm, random_state=0,
197-
positive_code=positive_code, positive_dict=positive_dict).fit(X)
220+
positive_code=positive_code, positive_dict=positive_dict,
221+
fit_algorithm='cd').fit(X)
222+
198223
code = dico.transform(X)
199224
if positive_dict:
200225
assert (dico.components_ >= 0).all()
@@ -205,7 +230,30 @@ def test_dict_learning_online_positivity(transform_algorithm,
205230
else:
206231
assert (code < 0).any()
207232

233+
234+
@pytest.mark.parametrize("positive_dict", [False, True])
235+
def test_minibatch_dictionary_learning_lars(positive_dict):
236+
n_components = 8
237+
238+
dico = MiniBatchDictionaryLearning(
239+
n_components, transform_algorithm="lars", random_state=0,
240+
positive_dict=positive_dict, fit_algorithm='cd').fit(X)
241+
242+
if positive_dict:
243+
assert (dico.components_ >= 0).all()
244+
else:
245+
assert (dico.components_ < 0).any()
246+
247+
248+
@pytest.mark.parametrize("positive_code", [False, True])
249+
@pytest.mark.parametrize("positive_dict", [False, True])
250+
def test_dict_learning_online_positivity(positive_code,
251+
positive_dict):
252+
rng = np.random.RandomState(0)
253+
n_components = 8
254+
208255
code, dictionary = dict_learning_online(X, n_components=n_components,
256+
method="cd",
209257
alpha=1, random_state=rng,
210258
positive_dict=positive_dict,
211259
positive_code=positive_code)
@@ -307,29 +355,34 @@ def test_sparse_encode_shapes():
307355
assert_equal(code.shape, (n_samples, n_components))
308356

309357

310-
# positive lars deprecated 0.22
311-
@pytest.mark.filterwarnings('ignore::DeprecationWarning')
312-
@pytest.mark.parametrize("positive", [
313-
False,
314-
True,
358+
@pytest.mark.parametrize("algo", [
359+
'lasso_lars',
360+
'lasso_cd',
361+
'threshold'
315362
])
316-
def test_sparse_encode_positivity(positive):
363+
@pytest.mark.parametrize("positive", [False, True])
364+
def test_sparse_encode_positivity(algo, positive):
317365
n_components = 12
318366
rng = np.random.RandomState(0)
319367
V = rng.randn(n_components, n_features) # random init
320368
V /= np.sum(V ** 2, axis=1)[:, np.newaxis]
321-
for algo in ('lasso_lars', 'lasso_cd', 'lars', 'threshold'):
322-
code = sparse_encode(X, V, algorithm=algo, positive=positive)
323-
if positive:
324-
assert (code >= 0).all()
325-
else:
326-
assert (code < 0).any()
369+
code = sparse_encode(X, V, algorithm=algo, positive=positive)
370+
if positive:
371+
assert (code >= 0).all()
372+
else:
373+
assert (code < 0).any()
327374

328-
try:
329-
sparse_encode(X, V, algorithm='omp', positive=positive)
330-
except ValueError:
331-
if not positive:
332-
raise
375+
376+
@pytest.mark.parametrize("algo", ['lars', 'omp'])
377+
def test_sparse_encode_unavailable_positivity(algo):
378+
n_components = 12
379+
rng = np.random.RandomState(0)
380+
V = rng.randn(n_components, n_features) # random init
381+
V /= np.sum(V ** 2, axis=1)[:, np.newaxis]
382+
err_msg = "Positive constraint not supported for '{}' coding method."
383+
err_msg = err_msg.format(algo)
384+
with pytest.raises(ValueError, match=err_msg):
385+
sparse_encode(X, V, algorithm=algo, positive=True)
333386

334387

335388
def test_sparse_encode_input():

sklearn/linear_model/least_angle.py

Lines changed: 10 additions & 27 deletions
Original file line numberDiff line numberDiff line change
@@ -397,10 +397,11 @@ def _lars_path_solver(X, y, Xy=None, Gram=None, n_samples=None, max_iter=500,
397397
398398
"""
399399
if method == 'lar' and positive:
400-
warnings.warn('positive option is broken for Least'
401-
' Angle Regression (LAR). Use method="lasso".'
402-
' This option will be removed in version 0.22.',
403-
DeprecationWarning)
400+
raise ValueError(
401+
"Positive constraint not supported for 'lar' "
402+
"coding method."
403+
)
404+
404405
n_samples = n_samples if n_samples is not None else y.size
405406

406407
if Xy is None:
@@ -804,14 +805,6 @@ class Lars(LinearModel, RegressorMixin, MultiOutputMixin):
804805
setting ``fit_path`` to ``False`` will lead to a speedup, especially
805806
with a small alpha.
806807
807-
positive : boolean (default=False)
808-
Restrict coefficients to be >= 0. Be aware that you might want to
809-
remove fit_intercept which is set True by default.
810-
811-
.. deprecated:: 0.20
812-
813-
The option is broken and deprecated. It will be removed in v0.22.
814-
815808
Attributes
816809
----------
817810
alphas_ : array, shape (n_alphas + 1,) | list of n_targets such arrays
@@ -844,7 +837,7 @@ class Lars(LinearModel, RegressorMixin, MultiOutputMixin):
844837
>>> reg.fit([[-1, 1], [0, 0], [1, 1]], [-1.1111, 0, -1.1111])
845838
... # doctest: +ELLIPSIS, +NORMALIZE_WHITESPACE
846839
Lars(copy_X=True, eps=..., fit_intercept=True, fit_path=True,
847-
n_nonzero_coefs=1, normalize=True, positive=False, precompute='auto',
840+
n_nonzero_coefs=1, normalize=True, precompute='auto',
848841
verbose=False)
849842
>>> print(reg.coef_) # doctest: +ELLIPSIS, +NORMALIZE_WHITESPACE
850843
[ 0. -1.11...]
@@ -856,17 +849,16 @@ class Lars(LinearModel, RegressorMixin, MultiOutputMixin):
856849
857850
"""
858851
method = 'lar'
852+
positive = False
859853

860854
def __init__(self, fit_intercept=True, verbose=False, normalize=True,
861855
precompute='auto', n_nonzero_coefs=500,
862-
eps=np.finfo(np.float).eps, copy_X=True, fit_path=True,
863-
positive=False):
856+
eps=np.finfo(np.float).eps, copy_X=True, fit_path=True):
864857
self.fit_intercept = fit_intercept
865858
self.verbose = verbose
866859
self.normalize = normalize
867860
self.precompute = precompute
868861
self.n_nonzero_coefs = n_nonzero_coefs
869-
self.positive = positive
870862
self.eps = eps
871863
self.copy_X = copy_X
872864
self.fit_path = fit_path
@@ -1303,13 +1295,6 @@ class LarsCV(Lars):
13031295
copy_X : boolean, optional, default True
13041296
If ``True``, X will be copied; else, it may be overwritten.
13051297
1306-
positive : boolean (default=False)
1307-
Restrict coefficients to be >= 0. Be aware that you might want to
1308-
remove fit_intercept which is set True by default.
1309-
1310-
.. deprecated:: 0.20
1311-
The option is broken and deprecated. It will be removed in v0.22.
1312-
13131298
Attributes
13141299
----------
13151300
coef_ : array, shape (n_features,)
@@ -1360,7 +1345,7 @@ class LarsCV(Lars):
13601345
def __init__(self, fit_intercept=True, verbose=False, max_iter=500,
13611346
normalize=True, precompute='auto', cv='warn',
13621347
max_n_alphas=1000, n_jobs=None, eps=np.finfo(np.float).eps,
1363-
copy_X=True, positive=False):
1348+
copy_X=True):
13641349
self.max_iter = max_iter
13651350
self.cv = cv
13661351
self.max_n_alphas = max_n_alphas
@@ -1369,8 +1354,7 @@ def __init__(self, fit_intercept=True, verbose=False, max_iter=500,
13691354
verbose=verbose, normalize=normalize,
13701355
precompute=precompute,
13711356
n_nonzero_coefs=500,
1372-
eps=eps, copy_X=copy_X, fit_path=True,
1373-
positive=positive)
1357+
eps=eps, copy_X=copy_X, fit_path=True)
13741358

13751359
def fit(self, X, y):
13761360
"""Fit the model using X, y as training data.
@@ -1683,7 +1667,6 @@ class LassoLarsIC(LassoLars):
16831667
As a consequence using LassoLarsIC only makes sense for problems where
16841668
a sparse solution is expected and/or reached.
16851669
1686-
16871670
Attributes
16881671
----------
16891672
coef_ : array, shape (n_features,)

sklearn/linear_model/tests/test_least_angle.py

Lines changed: 4 additions & 6 deletions
Original file line numberDiff line numberDiff line change
@@ -514,13 +514,11 @@ def test_lars_path_positive_constraint():
514514
# and all positive when positive=True
515515
# for method 'lar' (default) and lasso
516516

517-
# Once deprecation of LAR + positive option is done use these:
518-
# assert_raises(ValueError, linear_model.lars_path, diabetes['data'],
519-
# diabetes['target'], method='lar', positive=True)
520-
with pytest.warns(DeprecationWarning, match='broken'):
517+
err_msg = "Positive constraint not supported for 'lar' coding method."
518+
with pytest.raises(ValueError, match=err_msg):
521519
linear_model.lars_path(diabetes['data'], diabetes['target'],
522-
return_path=True, method='lar',
523-
positive=True)
520+
method='lar', positive=True)
521+
524522
method = 'lasso'
525523
_, _, coefs = \
526524
linear_model.lars_path(X, y, return_path=True, method=method,

0 commit comments

Comments
 (0)