Skip to content

Commit 1c2af45

Browse files
cajchristianChristian Jorgensen
andauthored
Adding argument for internal PCovC scaling
* Adding standardscaler calls * Fixing sample space pcovc scaling * Adjusting some params in the examples * oops wrong scaler * Touch up docs * Fix docs again * fix docs * One last docs fix * Adding tests * Fixed doctests * Changing default and adding warnings * Touching up docs * adding tests for kpcovc * Revert doctest changes * fixing examples * Doc fix --------- Co-authored-by: Christian Jorgensen <[email protected]>
1 parent 44407ee commit 1c2af45

File tree

7 files changed

+195
-8
lines changed

7 files changed

+195
-8
lines changed

examples/pcovc/KPCovC_Comparison.py

Lines changed: 5 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -33,6 +33,7 @@
3333

3434
random_state = 0
3535
n_components = 2
36+
scale_z = True
3637

3738
# %%
3839
#
@@ -85,7 +86,7 @@
8586
# Both PCA and PCovC fail to produce linearly separable latent space
8687
# maps. We will need a kernel method to effectively separate the moon classes.
8788

88-
mixing = 0.10
89+
mixing = 0.5
8990
alpha_d = 0.5
9091
alpha_p = 0.4
9192

@@ -95,6 +96,7 @@
9596
n_components=n_components,
9697
random_state=random_state,
9798
mixing=mixing,
99+
scale_z=scale_z,
98100
classifier=LinearSVC(),
99101
): "PCovC",
100102
}
@@ -138,6 +140,7 @@
138140
random_state=random_state,
139141
mixing=mixing,
140142
center=center,
143+
scale_z=scale_z,
141144
**kernel_params,
142145
): {"title": "Kernel PCovC", "eps": 2},
143146
}
@@ -220,6 +223,7 @@
220223
mixing=mixing,
221224
classifier=model,
222225
center=center,
226+
scale_z=scale_z,
223227
**models[model]["kernel_params"],
224228
)
225229
t_kpcovc_train = kpcovc.fit_transform(X_train_scaled, y_train)

examples/pcovc/KPCovC_Hyperparameters.py

Lines changed: 4 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -65,7 +65,8 @@
6565
fig, axs = plt.subplots(2, len(kernels), figsize=(len(kernels) * 4, 8))
6666

6767
center = True
68-
mixing = 0.10
68+
mixing = 0.5
69+
scale_z = True
6970

7071
for i, kernel in enumerate(kernels):
7172
kpca = KernelPCA(
@@ -83,6 +84,7 @@
8384
random_state=random_state,
8485
**kernel_params.get(kernel, {}),
8586
center=center,
87+
scale_z=scale_z,
8688
)
8789
t_kpcovc = kpcovc.fit_transform(X_scaled, y)
8890

@@ -118,7 +120,7 @@
118120
kpcovc = KernelPCovC(
119121
n_components=n_components,
120122
random_state=random_state,
121-
mixing=mixing,
123+
mixing=0.1,
122124
center=center,
123125
kernel="rbf",
124126
gamma=gamma,

examples/pcovc/PCovC_Hyperparameters.py

Lines changed: 2 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -77,6 +77,7 @@
7777
n_components=n_components,
7878
random_state=random_state,
7979
classifier=LogisticRegressionCV(),
80+
scale_z=True,
8081
)
8182

8283
pcovc.fit(X_scaled, y)
@@ -120,6 +121,7 @@
120121
n_components=n_components,
121122
random_state=random_state,
122123
classifier=model,
124+
scale_z=True,
123125
)
124126

125127
pcovc.fit(X_scaled, y)

src/skmatter/decomposition/_kernel_pcovc.py

Lines changed: 42 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -1,3 +1,4 @@
1+
import warnings
12
import numpy as np
23

34
from sklearn import clone
@@ -16,7 +17,7 @@
1617
from sklearn.linear_model._base import LinearClassifierMixin
1718
from sklearn.utils.multiclass import check_classification_targets, type_of_target
1819

19-
from skmatter.preprocessing import KernelNormalizer
20+
from skmatter.preprocessing import KernelNormalizer, StandardFlexibleScaler
2021
from skmatter.utils import check_cl_fit
2122
from skmatter.decomposition import _BaseKPCov
2223

@@ -86,6 +87,9 @@ class KernelPCovC(LinearClassifierMixin, _BaseKPCov):
8687
If None, ``sklearn.linear_model.LogisticRegression()``
8788
is used as the classifier.
8889
90+
scale_z: bool, default=False
91+
Whether to scale Z prior to eigendecomposition.
92+
8993
kernel : {"linear", "poly", "rbf", "sigmoid", "precomputed"} or callable, default="linear"
9094
Kernel.
9195
@@ -116,6 +120,14 @@ class KernelPCovC(LinearClassifierMixin, _BaseKPCov):
116120
and for matrix inversions.
117121
Must be of range [0.0, infinity).
118122
123+
z_mean_tol: float, default=1e-12
124+
Tolerance for the column means of Z.
125+
Must be of range [0.0, infinity).
126+
127+
z_var_tol: float, default=1.5
128+
Tolerance for the column variances of Z.
129+
Must be of range [0.0, infinity).
130+
119131
n_jobs : int, default=None
120132
The number of parallel jobs to run.
121133
:obj:`None` means 1 unless in a :obj:`joblib.parallel_backend` context.
@@ -167,14 +179,17 @@ class KernelPCovC(LinearClassifierMixin, _BaseKPCov):
167179
The data used to fit the model. This attribute is used to build kernels
168180
from new data.
169181
182+
scale_z: bool
183+
Whether Z is being scaled prior to eigendecomposition.
184+
170185
Examples
171186
--------
172187
>>> import numpy as np
173188
>>> from skmatter.decomposition import KernelPCovC
174189
>>> from sklearn.preprocessing import StandardScaler
175190
>>> X = np.array([[-2, 3, -1, 0], [2, 0, -3, 1], [3, 0, -1, 3], [2, -2, 1, 0]])
176191
>>> X = StandardScaler().fit_transform(X)
177-
>>> Y = np.array([[2], [0], [1], [2]])
192+
>>> Y = np.array([2, 0, 1, 2])
178193
>>> kpcovc = KernelPCovC(
179194
... mixing=0.1,
180195
... n_components=2,
@@ -200,6 +215,7 @@ def __init__(
200215
n_components=None,
201216
svd_solver="auto",
202217
classifier=None,
218+
scale_z=False,
203219
kernel="linear",
204220
gamma=None,
205221
degree=3,
@@ -208,6 +224,8 @@ def __init__(
208224
center=False,
209225
fit_inverse_transform=False,
210226
tol=1e-12,
227+
z_mean_tol=1e-12,
228+
z_var_tol=1.5,
211229
n_jobs=None,
212230
iterated_power="auto",
213231
random_state=None,
@@ -229,6 +247,9 @@ def __init__(
229247
fit_inverse_transform=fit_inverse_transform,
230248
)
231249
self.classifier = classifier
250+
self.scale_z = scale_z
251+
self.z_mean_tol = z_mean_tol
252+
self.z_var_tol = z_var_tol
232253

233254
def fit(self, X, Y, W=None):
234255
r"""Fit the model with X and Y.
@@ -323,6 +344,25 @@ def fit(self, X, Y, W=None):
323344
W = LogisticRegression().fit(K, Y).coef_.T
324345

325346
Z = K @ W
347+
if self.scale_z:
348+
Z = StandardFlexibleScaler().fit_transform(Z)
349+
350+
z_means_ = np.mean(Z, axis=0)
351+
z_vars_ = np.var(Z, axis=0)
352+
353+
if np.max(np.abs(z_means_)) > self.z_mean_tol:
354+
warnings.warn(
355+
"This class does not automatically center Z, and the column means "
356+
"of Z are greater than the supplied tolerance. We recommend scaling "
357+
"Z (and the weights) by setting `scale_z=True`."
358+
)
359+
360+
if np.max(z_vars_) > self.z_var_tol:
361+
warnings.warn(
362+
"This class does not automatically scale Z, and the column variances "
363+
"of Z are greater than the supplied tolerance. We recommend scaling "
364+
"Z (and the weights) by setting `scale_z=True`."
365+
)
326366

327367
self._fit(K, Z, W)
328368

src/skmatter/decomposition/_pcovc.py

Lines changed: 53 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -16,6 +16,8 @@
1616
from sklearn.utils.validation import check_is_fitted, validate_data
1717
from skmatter.decomposition import _BasePCov
1818
from skmatter.utils import check_cl_fit
19+
from skmatter.preprocessing import StandardFlexibleScaler
20+
import warnings
1921

2022

2123
class PCovC(LinearClassifierMixin, _BasePCov):
@@ -96,6 +98,14 @@ class PCovC(LinearClassifierMixin, _BasePCov):
9698
Tolerance for singular values computed by svd_solver == 'arpack'.
9799
Must be of range [0.0, infinity).
98100
101+
z_mean_tol: float, default=1e-12
102+
Tolerance for the column means of Z.
103+
Must be of range [0.0, infinity).
104+
105+
z_var_tol: float, default=1.5
106+
Tolerance for the column variances of Z.
107+
Must be of range [0.0, infinity).
108+
99109
space: {'feature', 'sample', 'auto'}, default='auto'
100110
whether to compute the PCovC in ``sample`` or ``feature`` space.
101111
The default is equal to ``sample`` when :math:`{n_{samples} < n_{features}}`
@@ -123,6 +133,9 @@ class PCovC(LinearClassifierMixin, _BasePCov):
123133
If None, ``sklearn.linear_model.LogisticRegression()``
124134
is used as the classifier.
125135
136+
scale_z: bool, default=False
137+
Whether to scale Z prior to eigendecomposition.
138+
126139
iterated_power : int or 'auto', default='auto'
127140
Number of iterations for the power method computed by
128141
svd_solver == 'randomized'.
@@ -143,6 +156,14 @@ class PCovC(LinearClassifierMixin, _BasePCov):
143156
Tolerance for singular values computed by svd_solver == 'arpack'.
144157
Must be of range [0.0, infinity).
145158
159+
z_mean_tol: float
160+
Tolerance for the column means of Z.
161+
Must be of range [0.0, infinity).
162+
163+
z_var_tol: float
164+
Tolerance for the column variances of Z.
165+
Must be of range [0.0, infinity).
166+
146167
space: {'feature', 'sample', 'auto'}, default='auto'
147168
whether to compute the PCovC in ``sample`` or ``feature`` space.
148169
The default is equal to ``sample`` when :math:`{n_{samples} < n_{features}}`
@@ -174,6 +195,9 @@ class PCovC(LinearClassifierMixin, _BasePCov):
174195
the projector, or weights, from the latent-space projection
175196
:math:`\mathbf{T}` to the class confidence scores :math:`\mathbf{Z}`
176197
198+
scale_z: bool
199+
Whether Z is being scaled prior to eigendecomposition
200+
177201
explained_variance_ : numpy.ndarray of shape (n_components,)
178202
The amount of variance explained by each of the selected components.
179203
Equal to n_components largest eigenvalues
@@ -208,8 +232,11 @@ def __init__(
208232
n_components=None,
209233
svd_solver="auto",
210234
tol=1e-12,
235+
z_mean_tol=1e-12,
236+
z_var_tol=1.5,
211237
space="auto",
212238
classifier=None,
239+
scale_z=False,
213240
iterated_power="auto",
214241
random_state=None,
215242
whiten=False,
@@ -225,6 +252,9 @@ def __init__(
225252
whiten=whiten,
226253
)
227254
self.classifier = classifier
255+
self.scale_z = scale_z
256+
self.z_mean_tol = z_mean_tol
257+
self.z_var_tol = z_var_tol
228258

229259
def fit(self, X, Y, W=None):
230260
r"""Fit the model with X and Y.
@@ -291,7 +321,7 @@ def fit(self, X, Y, W=None):
291321
classifier = self.classifier
292322

293323
self.z_classifier_ = check_cl_fit(classifier, X, Y)
294-
W = self.z_classifier_.coef_.T
324+
W = self.z_classifier_.coef_.T.copy()
295325

296326
else:
297327
# If precomputed, use default classifier to predict Y from T
@@ -301,6 +331,28 @@ def fit(self, X, Y, W=None):
301331

302332
Z = X @ W
303333

334+
if self.scale_z:
335+
z_scaler = StandardFlexibleScaler().fit(Z)
336+
Z = z_scaler.transform(Z)
337+
W /= z_scaler.scale_.reshape(1, -1)
338+
339+
z_means_ = np.mean(Z, axis=0)
340+
z_vars_ = np.var(Z, axis=0)
341+
342+
if np.max(np.abs(z_means_)) > self.z_mean_tol:
343+
warnings.warn(
344+
"This class does not automatically center Z, and the column means "
345+
"of Z are greater than the supplied tolerance. We recommend scaling "
346+
"Z (and the weights) by setting `scale_z=True`."
347+
)
348+
349+
if np.max(z_vars_) > self.z_var_tol:
350+
warnings.warn(
351+
"This class does not automatically scale Z, and the column variances "
352+
"of Z are greater than the supplied tolerance. We recommend scaling "
353+
"Z (and the weights) by setting `scale_z=True`."
354+
)
355+
304356
if self.space_ == "feature":
305357
self._fit_feature_space(X, Y, Z)
306358
else:

tests/test_kernel_pcovc.py

Lines changed: 41 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1,4 +1,5 @@
11
import unittest
2+
import warnings
23

34
import numpy as np
45
from sklearn import exceptions
@@ -34,10 +35,12 @@ def __init__(self, *args, **kwargs):
3435
lambda mixing=0.5,
3536
classifier=LogisticRegression(),
3637
n_components=4,
38+
scale_z=True,
3739
**kwargs: KernelPCovC(
3840
mixing=mixing,
3941
classifier=classifier,
4042
n_components=n_components,
43+
scale_z=scale_z,
4144
svd_solver=kwargs.pop("svd_solver", "full"),
4245
**kwargs,
4346
)
@@ -327,6 +330,44 @@ def test_precomputed_classification(self):
327330
self.assertTrue(np.linalg.norm(t3 - t2) < self.error_tol)
328331
self.assertTrue(np.linalg.norm(t3 - t1) < self.error_tol)
329332

333+
def test_scale_z_parameter(self):
334+
"""Check that changing scale_z changes the eigendecomposition."""
335+
kpcovc_scaled = self.model(scale_z=True)
336+
kpcovc_scaled.fit(self.X, self.Y)
337+
338+
kpcovc_unscaled = self.model(scale_z=False)
339+
kpcovc_unscaled.fit(self.X, self.Y)
340+
assert not np.allclose(kpcovc_scaled.pkt_, kpcovc_unscaled.pkt_)
341+
342+
def test_z_scaling(self):
343+
"""
344+
Check that KPCovC raises a warning if Z is not of scale, and does not
345+
if it is.
346+
"""
347+
kpcovc = self.model(n_components=2, scale_z=True)
348+
349+
with warnings.catch_warnings():
350+
kpcovc.fit(self.X, self.Y)
351+
warnings.simplefilter("error")
352+
self.assertEqual(1 + 1, 2)
353+
354+
kpcovc = self.model(n_components=2, scale_z=False, z_mean_tol=0, z_var_tol=0)
355+
356+
with warnings.catch_warnings(record=True) as w:
357+
kpcovc.fit(self.X, self.Y)
358+
self.assertEqual(
359+
str(w[0].message),
360+
"This class does not automatically center Z, and the column means "
361+
"of Z are greater than the supplied tolerance. We recommend scaling "
362+
"Z (and the weights) by setting `scale_z=True`.",
363+
)
364+
self.assertEqual(
365+
str(w[1].message),
366+
"This class does not automatically scale Z, and the column variances "
367+
"of Z are greater than the supplied tolerance. We recommend scaling "
368+
"Z (and the weights) by setting `scale_z=True`.",
369+
)
370+
330371

331372
class KernelTests(KernelPCovCBaseTest):
332373
def test_kernel_types(self):

0 commit comments

Comments
 (0)