Skip to content

Commit ae950da

Browse files
committed
Implement complex lasso in progress ......
1 parent 0fa9ecf commit ae950da

File tree

2 files changed

+176
-0
lines changed

2 files changed

+176
-0
lines changed

sklearn/linear_model/cd_fast.pyx

Lines changed: 92 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -874,3 +874,95 @@ def enet_coordinate_descent_multi_task(floating[::1, :] W, floating l1_reg,
874874
break
875875

876876
return np.asarray(W), gap, tol, n_iter + 1
877+
878+
@cython.boundscheck(False)
879+
@cython.wraparound(False)
880+
@cython.cdivision(True)
881+
def enet_coordinate_descent_complex(floating[::1, :] W, floating l1_reg,
882+
floating l2_reg,
883+
floating[::1, :] Xr,
884+
floating[::1, :] Xi,
885+
floating[::1, :] Y,
886+
int max_iter, floating tol, object rng,
887+
bint random=0):
888+
"""Cython version of the coordinate descent algorithm
889+
for Elastic-Net mult-task regression in complex domain
890+
891+
"""
892+
893+
# fused types version of BLAS functions
894+
if floating is float:
895+
dtype = np.float32
896+
gemv = sgemv
897+
dot = sdot
898+
copy = scopy
899+
else:
900+
dtype = np.float64
901+
gemv = dgemv
902+
dot = ddot
903+
copy = dcopy
904+
905+
# get the data information into easy vars
906+
cdef unsigned int n_samples = Xr.shape[0]
907+
cdef unsigned int n_features = Xr.shape[1]
908+
909+
# initial value of the residuals
910+
cdef floating[::1] Rr = np.empty(n_samples, dtype=dtype)
911+
cdef floating[::1] Ri = np.empty(n_samples, dtype=dtype)
912+
913+
cdef floating[:] w_ii = np.zeros(2, dtype=dtype)
914+
cdef unsigned int ii
915+
cdef unsigned int jj
916+
cdef unsigned int n_iter = 0
917+
cdef unsigned int f_iter
918+
cdef UINT32_t rand_r_state_seed = rng.randint(0, RAND_R_MAX)
919+
cdef UINT32_t* rand_r_state = &rand_r_state_seed
920+
921+
cdef floating* W_ptr = &W[0, 0]
922+
cdef floating* Y_ptr = &Y[0, 0]
923+
924+
if l1_reg == 0:
925+
warnings.warn("Coordinate descent with l1_reg=0 may lead to unexpected"
926+
" results and is discouraged.")
927+
928+
with nogil:
929+
# Compute Rr and Ri: real and imaginary parts of the residual
930+
# real part: Yr - np.dot(Xr, Wr) + np.dot(Xi, Wi)
931+
copy(n_samples, Y_ptr, 1, &Rr[0], 1)
932+
gemv(CblasColMajor, CblasNoTrans,
933+
n_samples, n_features, -1.0, &Xr[0, 0], n_samples,
934+
W_ptr, 2, 1.0, &Rr[0], 1)
935+
gemv(CblasColMajor, CblasNoTrans,
936+
n_samples, n_features, 1.0, &Xi[0, 0], n_samples,
937+
W_ptr + 1, 2, 1.0, &Rr[0], 1)
938+
939+
# imaginary part:
940+
# real part: Yr - np.dot(Xr, Wi) - np.dot(Xi, Wr)
941+
copy(n_samples, Y_ptr + n_samples, 1, &Ri[0], 1)
942+
gemv(CblasColMajor, CblasNoTrans,
943+
n_samples, n_features, -1.0, &Xr[0, 0], n_samples,
944+
W_ptr + 1, 2, 1.0, &Ri[0], 1)
945+
gemv(CblasColMajor, CblasNoTrans,
946+
n_samples, n_features, -1.0, &Xi[0, 0], n_samples,
947+
W_ptr, 2, 1.0, &Ri[0], 1)
948+
949+
# tol = tol * linalg.norm(Y, ord='fro') ** 2
950+
tol = tol * dot(n_samples * 2, Y_ptr, 1, Y_ptr, 1)
951+
952+
for n_iter in range(max_iter):
953+
for f_iter in range(n_features): # Loop over coordinates
954+
# select a coordinate
955+
if random:
956+
ii = rand_int(n_features, rand_r_state)
957+
else:
958+
ii = f_iter
959+
960+
# w_ii = W[:, ii] # Store previous value
961+
w_ii[0] = W[0, ii]
962+
w_ii[1] = W[1, ii]
963+
964+
# if np.sum(w_ii ** 2) != 0.0: # can do better
965+
if w_ii[0] != 0.0 or w_ii[1] != 0.0:
966+
# Remove contributions of w_ii from R
967+
968+
# prepare for the soft-thresholding

sklearn/linear_model/coordinate_descent.py

Lines changed: 84 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -2260,3 +2260,87 @@ def __init__(self, eps=1e-3, n_alphas=100, alphas=None, fit_intercept=True,
22602260
max_iter=max_iter, tol=tol, copy_X=copy_X,
22612261
cv=cv, verbose=verbose, n_jobs=n_jobs, random_state=random_state,
22622262
selection=selection)
2263+
2264+
###############################################################################
2265+
# Complex ElasticNet and Lasso models
2266+
2267+
class ComplexElasticNet(ElasticNet):
2268+
"""ElasticNet model in complex domain trained with L1/L2 mixed-norm as regularizer
2269+
"""
2270+
2271+
def __init__(self, alpha=1.0, l1_ratio=0.5, max_iter=1000,
2272+
tol=1e-4, warm_start=False, random_state=None,
2273+
selection='cyclic'):
2274+
self.l1_ratio = l1_ratio
2275+
self.alpha = alpha
2276+
self.max_iter = max_iter
2277+
self.tol = tol
2278+
self.warm_start = warm_start
2279+
self.random_state = random_state
2280+
self.selection = selection
2281+
2282+
def fit(self, Xr, Xi, Y):
2283+
"""Fit MultiTaskElasticNet model with coordinate descent
2284+
2285+
Parameters
2286+
-----------
2287+
Xr : ndarray, shape (n_samples, n_features)
2288+
real part of the X matrix
2289+
Xi : ndarray, shape (n_samples, n_features)
2290+
imaginary part of the X matrix
2291+
Y : ndarray, shape (n_samples, 2)
2292+
Target. Will be cast to X's dtype if necessary
2293+
Y[:,0] is the real part and Y[:,1] is the imaginary part
2294+
"""
2295+
Xr = check_array(Xr, dtype=[np.float64, np.float32], order='F', copy=False)
2296+
Xi = check_array(Xi, dtype=[np.float64, np.float32], order='F', copy=False)
2297+
Y = check_array(Y, dtype=X.dtype.type, order='F', copy=False)
2298+
2299+
n_samples, n_features = X.shape
2300+
2301+
if n_samples != Y.shape[0]:
2302+
raise ValueError("X and Y have inconsistent dimensions (%d != %d)"
2303+
% (n_samples, y.shape[0]))
2304+
if Y.shape[1] != 2:
2305+
raise ValueError("Y should have %d columns, but two are expected (real and imaginary part)"
2306+
% y.shape[1])
2307+
2308+
if not self.warm_start or self.coef_ is None:
2309+
self.coef_ = np.zeros((2, n_features), dtype=X.dtype.type, order='F')
2310+
2311+
l1_reg = self.alpha * self.l1_ratio * n_samples
2312+
l2_reg = self.alpha * (1.0 - self.l1_ratio) * n_samples
2313+
2314+
self.coef_ = np.asfortranarray(self.coef_) # coef contiguous in memory
2315+
2316+
if self.selection not in ['random', 'cyclic']:
2317+
raise ValueError("selection should be either random or cyclic.")
2318+
random = (self.selection == 'random')
2319+
2320+
# now call coordinate descient
2321+
self.coef_, self.dual_gap_, self.eps_, self.n_iter_ = \
2322+
cd_fast.enet_coordinate_descent_complex(
2323+
self.coef_, l1_reg, l2_reg, Xr, Xi, Y, self.max_iter, self.tol,
2324+
check_random_state(self.random_state), random)
2325+
2326+
if self.dual_gap_ > self.eps_:
2327+
warnings.warn('Objective did not converge, you might want'
2328+
' to increase the number of iterations',
2329+
ConvergenceWarning)
2330+
2331+
# return self for chaining fit and predict calls
2332+
return self
2333+
2334+
class ComplexLasso(ComplexElasticNet):
2335+
2336+
def __init__(self, alpha=1.0, max_iter=1000, tol=1e-4,
2337+
warm_start=False, random_state=None, selection='cyclic'):
2338+
self.l1_ratio = 1.0
2339+
self.alpha = alpha
2340+
self.max_iter = max_iter
2341+
self.tol = tol
2342+
self.warm_start = warm_start
2343+
self.random_state = random_state
2344+
self.selection = selection
2345+
2346+
###############################################################################

0 commit comments

Comments
 (0)