CSC#

class brainevent.CSC(data, indices=None, indptr=None, *, shape, backend=None, buffers=None)#

Event-driven and Unit-aware Compressed Sparse Column (CSC) matrix.

This class represents a sparse matrix in CSC format, which is efficient for column-wise operations. It is compatible with JAX’s tree utilities and supports unit-aware computations.

The class supports arithmetic with scalars and dense arrays, plus sparse-dense matrix multiplication via @. Sparse-sparse operations are limited.

data#

Array of the non-zero values in the matrix.

Type:

Data

indices#

Array of row indices for the non-zero values.

Type:

jax.Array

indptr#

Array of column pointers indicating where each column starts in the data and indices arrays.

Type:

jax.Array

shape#

The shape of the matrix as (rows, columns).

Type:

tuple[int, int]

nse#

Number of stored elements (non-zero entries).

Type:

int

dtype#

Data type of the matrix values.

Type:

dtype

Notes

In CSC format a matrix of shape (m, n) is stored as three arrays:

  • indptr of length n + 1 – the j-th column occupies entries indptr[j] to indptr[j+1] in the data and indices arrays.

  • indices – row indices of the stored elements.

  • data – the corresponding non-zero values.

Internally, CSC operations are implemented by treating the underlying arrays as a CSR matrix with transposed shape and applying the appropriate transpose flags to the CSR kernels.

Examples

import jax.numpy as jnp
import brainevent

data    = jnp.array([1.0, 2.0, 3.0])
indices = jnp.array([0, 2, 1])
indptr  = jnp.array([0, 1, 2, 3])
csc     = brainevent.CSC((data, indices, indptr), shape=(3, 3))

# Sparse-dense matrix-vector product
x = jnp.ones(3)
y = csc @ x

See also

CSR

Compressed Sparse Row format.

apply(fn)[source]#

Apply a unary function to the stored data values.

Creates a new CSC matrix with fn(self.data) while preserving the sparsity structure (indices, indptr, shape, and cached diagonal positions).

Parameters:

fn (callable) – A function that accepts a single array argument and returns an array of the same shape. The dtype and unit may differ from the input.

Returns:

A new CSC matrix with transformed data.

Return type:

CSC

Examples

squared = csc.apply(lambda x: x ** 2)
build_weight_indices()[source]#

Return a copy of this CSC with the weight indices eagerly cached.

Builds the row-major (CSR-like) structure and permutation used by the event @ CSC direction (see _weight_indices()) and stores it in the 'csr' buffer of the returned matrix. The underlying data, indices, and indptr arrays are shared (not copied).

Returns:

A new CSC matrix sharing this matrix’s arrays, with the 'csr' weight-index buffer populated.

Return type:

CSC

See also

CSC._weight_indices

Lazy builder/accessor for the same triple.

CSC.fromdense

Accepts precompute_weight_indices=True to call this.

classmethod fromdense(mat, *, nse=None, index_dtype=<class 'jax.numpy.int32'>, backend=None, precompute_weight_indices=False)[source]#

Create a CSC (Compressed Sparse Column) matrix from a dense matrix.

This method converts a dense matrix to CSC format, which is an efficient storage format for sparse matrices.

Parameters:
  • mat (array_like) – The dense matrix to be converted to CSC format.

  • nse (int) – The number of non-zero elements in the matrix. If None, it will be calculated from the input matrix.

  • index_dtype (dtype, optional) – The data type to be used for index arrays (default is jnp.int32).

  • backend (str | None) – Compute backend to attach to the matrix. Default None.

  • precompute_weight_indices (bool) – If True, eagerly build and cache the row-major (CSR-like) weight indices used by the unfavorable event @ CSC direction (see build_weight_indices()). If False (default), the indices are built lazily on first use. Default False.

Returns:

A new CSC matrix instance created from the input dense matrix.

Return type:

CSC

See also

build_weight_indices

Eagerly build the cached weight indices.

CSC._weight_indices

Lazily build/return the cached weight indices.

Examples

import jax.numpy as jnp
import brainevent

dense = jnp.array([[1.0, 0.0], [0.0, 2.0]])
csc = brainevent.CSC.fromdense(dense)
slice_rows(index)[source]#

Return W[rows, :] as a new CSC (outside jax.jit).

Builds the CSR arrays of W[rows, :] through the cached CSR-of-W view, then converts to CSC. The output non-zero count is data-dependent, so index must be concrete.

Parameters:

index (int, list, tuple, array, or slice) – Row selector along axis 0.

Returns:

Sparse sub-matrix of shape (len(rows), n_cols).

Return type:

CSC

solve(b, tol=1e-06, reorder=1)[source]#

Solve the linear system A x = b where A is this CSC matrix.

Delegates to the CSR solver by transposing the matrix.

Parameters:
  • b (Array | Quantity) – Right-hand side vector. Its first dimension must equal self.shape[0].

  • tol (float, optional) – Tolerance for singularity detection. Defaults to 1e-6.

  • reorder (int, optional) – Fill-reducing reordering scheme: 0 for no reordering, 1 for symrcm, 2 for symamd, 3 for csrmetisnd. Defaults to 1.

Returns:

Solution vector x satisfying A x = b.

Return type:

Array | Quantity

Raises:

AssertionError – If b.shape[0] != self.shape[0].

Examples

x = csc.solve(b)
tocoo()[source]#

Convert to coordinate (COO) format.

Returns:

The same logical matrix in COO format, shape unchanged. A homogeneous (size-1) value is broadcast to one entry per stored element.

Return type:

COO

See also

tocsr

Re-encode the same logical matrix row-major.

tocsc

Identity conversion.

tocsc()[source]#

Return this matrix in CSC format (a no-op that returns self).

Provided for a uniform conversion interface across data representations; CSC is already column-compressed.

Returns:

self, unchanged.

Return type:

CSC

See also

tocsr

Re-encode the same logical matrix row-major.

tocoo

Convert to coordinate format.

tocsr()[source]#

Re-encode the same logical matrix in CSR format.

Unlike transpose() (which reinterprets the arrays as W.T with swapped shape), tocsr returns a CSR describing the same matrix W with the same shape – the entries are resorted into row-major order.

Returns:

The same logical matrix in CSR format, shape unchanged.

Return type:

CSR

See also

tocsc

Identity conversion.

transpose

Logical transpose (swaps shape).

todense()[source]#

Convert the CSC matrix to a dense matrix.

Transposes the underlying CSR-style storage, converts to dense, and transposes back.

Returns:

A dense matrix of shape self.shape containing all the values (including zeros) of the sparse matrix.

Return type:

Array | Quantity

Examples

dense = csc.todense()
transpose(axes=None)[source]#

Transpose the CSC matrix.

Returns the transpose as a CSR matrix. Because the transpose of a CSC matrix is a CSR matrix with the same underlying arrays, this operation is essentially free.

Parameters:

axes (None) – Must be None. Included for API compatibility with NumPy.

Returns:

The transpose of the CSC matrix as a CSR (Compressed Sparse Row) matrix.

Return type:

CSR

Raises:

AssertionError – If axes is not None.

Examples

csr = csc.transpose()
# or equivalently:
csr = csc.T
update_on_post(pre_trace, post_spike, w_min=None, w_max=None)[source]#

Apply a postsynaptic-spike-triggered STDP update, returning a new CSC.

Iterating by postsynaptic spike is the favorable direction for CSC, so this streams directly over the stored arrays (no permutation) via brainevent.update_csr_on_binary_pre() on the transposed shape. For each firing postsynaptic neuron j every stored synapse is updated W[i, j] <- clip(W[i, j] + pre_trace[i], w_min, w_max).

Parameters:
  • pre_trace (jax.Array or Quantity) – Presynaptic eligibility trace, shape (shape[0],).

  • post_spike (jax.Array) – Binary/boolean postsynaptic spikes, shape (shape[1],).

  • w_min (optional) – Clipping bounds; None disables the corresponding bound.

  • w_max (optional) – Clipping bounds; None disables the corresponding bound.

Returns:

A new CSC matrix with updated data and identical structure.

Return type:

CSC

See also

update_on_pre

Presynaptic-spike-triggered counterpart.

brainevent.update_csc_on_binary_post

Equivalent module function.

update_on_pre(pre_spike, post_trace, w_min=None, w_max=None)[source]#

Apply a presynaptic-spike-triggered STDP update, returning a new CSC.

Iterating by presynaptic spike is the unfavorable direction for CSC, so this reuses the cached row-major weight indices (_weight_indices()) and routes through brainevent.update_csr_on_binary_post(), scattering updates back into canonical CSC order. For each firing presynaptic neuron i every stored synapse is updated W[i, j] <- clip(W[i, j] + post_trace[j], w_min, w_max).

Parameters:
  • pre_spike (jax.Array) – Binary/boolean presynaptic spikes, shape (shape[0],).

  • post_trace (jax.Array or Quantity) – Postsynaptic eligibility trace, shape (shape[1],).

  • w_min (optional) – Clipping bounds; None disables the corresponding bound.

  • w_max (optional) – Clipping bounds; None disables the corresponding bound.

Returns:

A new CSC matrix with updated data and identical structure.

Return type:

CSC

See also

update_on_post

Postsynaptic-spike-triggered counterpart.

brainevent.update_csc_on_binary_pre

Equivalent module function.

with_data(data)[source]#

Create a new CSC matrix with updated data while keeping the same structure.

This method creates a new CSC matrix instance with the provided data, maintaining the original indices, indptr, and shape.

Parameters:

data (Array | ndarray | Quantity | Number) – The new data array to replace the existing data in the CSC matrix. It must have the same shape, dtype, and unit as the original data.

Returns:

A new CSC matrix instance with updated data and the same structure as the original.

Return type:

CSC

Raises:

AssertionError – If the shape, dtype, or unit of the new data doesn’t match the original data.

Examples

new_data = jnp.array([10.0, 20.0, 30.0])
new_csc = csc.with_data(new_data)
yw_to_w(y_dim_arr, w_dim_arr)[source]#

Compute a sparse transformation from y-w space to w space.

Performs a specialised sparse matrix-vector product optimised for event-driven neural simulations. The CSC storage is treated as a transposed CSR matrix, so this method calls the CSR kernel with shape=self.shape[::-1] and transpose=True.

Parameters:
  • y_dim_arr (Array | ndarray | Quantity) – Values in the target (post-synaptic) dimension.

  • w_dim_arr (Array | ndarray | Quantity) – Per-synapse weight values.

Returns:

Accumulated result, preserving physical units when present.

Return type:

Array | Quantity

See also

yw_to_w_transposed

The transposed (adjoint) variant.

Notes

Internally calls csrmv_yw2y with transpose=True and reversed shape to account for the column-oriented storage format.

yw_to_w_transposed(y_dim_arr, w_dim_arr)[source]#

Compute the transposed sparse transformation from y-w space to w space.

This is the adjoint of yw_to_w(), useful for back-propagation or adjoint computations in event-driven neural simulations. Uses transpose=False with the reversed shape to compute the appropriate transposed operation for CSC storage.

Parameters:
  • y_dim_arr (Array | ndarray | Quantity) – Values in the target (post-synaptic) dimension.

  • w_dim_arr (Array | ndarray | Quantity) – Per-synapse weight values.

Returns:

Accumulated result of the transposed operation, preserving physical units when present.

Return type:

Array | Quantity

See also

yw_to_w

The forward (non-transposed) variant.

Notes

Internally calls csrmv_yw2y with transpose=False and reversed shape.