Skip to content
New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

[WIP] Add Numpy Backend #87

Closed
wants to merge 54 commits into from
Closed
Show file tree
Hide file tree
Changes from all commits
Commits
Show all changes
54 commits
Select commit Hold shift + click to select a range
648c4ad
chore: adding numpy backend
ariG23498 May 4, 2023
d9627f3
creview comments
ariG23498 May 4, 2023
df8ec30
review comments
ariG23498 May 4, 2023
2309cb6
Merge branch 'aritra-np-backend' of https://github.com/keras-team/ker…
ariG23498 May 4, 2023
104571b
Merge branch 'main' into aritra-np-backend
ariG23498 May 5, 2023
2b4dbde
chore: adding math
ariG23498 May 5, 2023
99bca4a
Merge branch 'main' into aritra-np-backend
ariG23498 May 8, 2023
155a5b6
chore: adding random module
ariG23498 May 8, 2023
592fd6c
chore: adding ranndom in init
ariG23498 May 10, 2023
e97b21e
Merge branch 'main' into aritra-np-backend
ariG23498 May 10, 2023
8ce4450
review comments
ariG23498 May 10, 2023
65c8076
chore: adding numpy and nn for numpy backend
ariG23498 May 10, 2023
8f5dd4d
chore: adding generic pool, max, and average pool
ariG23498 May 18, 2023
8f22906
chore: adding the conv ops
ariG23498 May 18, 2023
b2c3184
Merge branch 'main' into aritra-np-backend
ariG23498 Jun 5, 2023
97f4e9e
chore: reformat code and using jax for conv and pool
ariG23498 Jun 5, 2023
d650587
chore: added self value
ariG23498 Jun 6, 2023
408f3e8
chore: activation tests pass
ariG23498 Jun 7, 2023
a846b34
chore: adding post build method
ariG23498 Jun 9, 2023
3e0283a
Merge branch 'main' into aritra-np-backend
ariG23498 Jun 9, 2023
5de62ef
Merge branch 'main' into aritra-np-backend
ariG23498 Jun 19, 2023
7143f06
chore: adding necessaity methods to the numpy trainer
ariG23498 Jun 19, 2023
5b4d800
chore: fixing utils test
ariG23498 Jun 19, 2023
4586e33
chore: fixing losses test suite
ariG23498 Jun 21, 2023
8373103
Merge branch 'main' into aritra-np-backend
ariG23498 Jun 21, 2023
82a1e5c
chore: fix backend tests
ariG23498 Jun 21, 2023
295e0e4
chore: fixing initializers test
ariG23498 Jun 21, 2023
dad6c9d
Merge branch 'main' into aritra-np-backend
ariG23498 Jun 22, 2023
55b5e09
chore: fixing accuracy metrics test
ariG23498 Jun 22, 2023
c66abaf
chore: fixing ops test
ariG23498 Jun 22, 2023
ca4869b
chore: review comments
ariG23498 Jun 27, 2023
926e169
chore: init with image and fixing random tests
ariG23498 Jun 27, 2023
23cd5b3
chore: skipping random seed set for numpy backend
ariG23498 Jun 27, 2023
52f8677
Merge branch 'main' into aritra-np-backend
ariG23498 Jun 30, 2023
e013d7d
chore: adding single resize image method
ariG23498 Jun 30, 2023
f6073cd
Merge branch 'main' into aritra-np-backend
ariG23498 Jul 7, 2023
17a5dda
chore: skipping tests for applications and layers
ariG23498 Jul 7, 2023
512c441
chore: skipping tests for models
ariG23498 Jul 7, 2023
f6f6442
chore: skipping testsor saving
ariG23498 Jul 7, 2023
bd38a79
chore: skipping tests for trainers
ariG23498 Jul 7, 2023
e29a54e
chore:ixing one hot
ariG23498 Jul 8, 2023
5694b25
Merge branch 'main' into aritra-np-backend
ariG23498 Jul 8, 2023
9d639cd
chore: fixing vmap in numpy and metrics test
ariG23498 Jul 8, 2023
6c8293b
chore: adding a wrapper to numpy sum, started fixing layer tests
ariG23498 Jul 8, 2023
f007fe0
fix: is_tensor now accepts numpy scalars
ariG23498 Jul 10, 2023
95abe6e
chore: adding draw seed
ariG23498 Jul 11, 2023
3547edc
Merge branch 'main' into aritra-np-backend
ariG23498 Jul 11, 2023
5bedccf
fix: warn message for numpy masking
ariG23498 Jul 11, 2023
f103ae0
fix: checking whether kernel are tensors
ariG23498 Jul 11, 2023
fe6bcf6
chore: adding rnn
ariG23498 Jul 11, 2023
360d913
chore: adding dynamic backend for numpy
ariG23498 Jul 11, 2023
b78500c
fix: axis cannot be None for normalize
ariG23498 Jul 11, 2023
13f256c
chore: adding jax resize for numpy image
ariG23498 Jul 11, 2023
9a88fa7
Merge branch 'main' into aritra-np-backend
ariG23498 Jul 11, 2023
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
4 changes: 4 additions & 0 deletions keras_core/applications/applications_test.py
Original file line number Diff line number Diff line change
Expand Up @@ -107,6 +107,10 @@ def _get_elephant(target_size):
os.environ.get("SKIP_APPLICATIONS_TESTS"),
reason="Env variable set to skip.",
)
@pytest.mark.skipif(
backend.backend() == "numpy",
reason="Trainer not implemented from NumPy backend.",
)
class ApplicationsTest(testing.TestCase, parameterized.TestCase):
@parameterized.named_parameters(MODEL_LIST)
def test_application_notop_variable_input_channels(self, app, last_dim, _):
Expand Down
6 changes: 6 additions & 0 deletions keras_core/applications/imagenet_utils_test.py
Original file line number Diff line number Diff line change
@@ -1,7 +1,9 @@
import numpy as np
import pytest
from absl.testing import parameterized

import keras_core as keras
from keras_core import backend
from keras_core import testing
from keras_core.applications import imagenet_utils as utils
from keras_core.mixed_precision import set_dtype_policy
Expand Down Expand Up @@ -74,6 +76,10 @@ def test_preprocess_input(self):
{"testcase_name": "mode_caffe", "mode": "caffe"},
]
)
@pytest.mark.skipif(
backend.backend() == "numpy",
reason="Trainer not implemented from NumPy backend.",
)
def test_preprocess_input_symbolic(self, mode):
# Test image batch
x = np.random.uniform(0, 255, (2, 10, 10, 3))
Expand Down
3 changes: 3 additions & 0 deletions keras_core/backend/__init__.py
Original file line number Diff line number Diff line change
Expand Up @@ -37,5 +37,8 @@
elif backend() == "torch":
print_msg("Using PyTorch backend.")
from keras_core.backend.torch import * # noqa: F403
elif backend() == "numpy":
print_msg("Using NumPy backend.")
from keras_core.backend.numpy import * # noqa: F403
else:
raise ValueError(f"Unable to import backend : {backend()}")
20 changes: 20 additions & 0 deletions keras_core/backend/numpy/__init__.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,20 @@
from keras_core.backend.numpy import core
from keras_core.backend.numpy import image
from keras_core.backend.numpy import math
from keras_core.backend.numpy import nn
from keras_core.backend.numpy import numpy
from keras_core.backend.numpy import random
from keras_core.backend.numpy.core import DYNAMIC_SHAPES_OK
from keras_core.backend.numpy.core import Variable
from keras_core.backend.numpy.core import cast
from keras_core.backend.numpy.core import compute_output_spec
from keras_core.backend.numpy.core import cond
from keras_core.backend.numpy.core import convert_to_numpy
from keras_core.backend.numpy.core import convert_to_tensor
from keras_core.backend.numpy.core import is_tensor
from keras_core.backend.numpy.core import name_scope
from keras_core.backend.numpy.core import shape
from keras_core.backend.numpy.core import vectorized_map
from keras_core.backend.numpy.rnn import gru
from keras_core.backend.numpy.rnn import lstm
from keras_core.backend.numpy.rnn import rnn
214 changes: 214 additions & 0 deletions keras_core/backend/numpy/core.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,214 @@
from contextlib import nullcontext

import numpy as np
from tensorflow import nest

from keras_core.backend.common import KerasVariable
from keras_core.backend.common import standardize_dtype
from keras_core.backend.common.keras_tensor import KerasTensor
from keras_core.backend.common.stateless_scope import StatelessScope

DYNAMIC_SHAPES_OK = True


class Variable(KerasVariable):
def _initialize(self, value):
self._value = np.array(value, dtype=self._dtype)

def _direct_assign(self, value):
self._value = np.array(value, dtype=self._dtype)

def _convert_to_tensor(self, value, dtype=None):
return convert_to_tensor(value, dtype=dtype)

# Overload native accessor.
def __array__(self):
return self.value


def convert_to_tensor(x, dtype=None):
if dtype is not None:
dtype = standardize_dtype(dtype)
if isinstance(x, Variable):
if dtype and dtype != x.dtype:
return x.value.astype(dtype)
return x.value
return np.array(x, dtype=dtype)


def convert_to_numpy(x):
return np.array(x)


def is_tensor(x):
if isinstance(x, (np.generic, np.ndarray)):
return True
return False


def shape(x):
# This will work as long as we disallow
# dynamic shapes in NumPy.
return x.shape


def cast(x, dtype):
return convert_to_tensor(x, dtype=dtype)


def cond(pred, true_fn, false_fn):
if pred:
return true_fn()
return false_fn()


def name_scope(name):
# There is no need for a named context for NumPy.
return nullcontext()


def vectorized_map(function, elements):
if len(elements) == 1:
return function(elements)
else:
batch_size = elements[0].shape[0]
output_store = list()
for index in range(batch_size):
output_store.append(function([x[index] for x in elements]))
return np.stack(output_store)


# Shape / dtype inference util
def compute_output_spec(fn, *args, **kwargs):
with StatelessScope():

def has_none_shape(x):
if isinstance(x, KerasTensor):
return None in x.shape
return False

none_in_shape = any(map(has_none_shape, nest.flatten((args, kwargs))))

def convert_keras_tensor_to_numpy(x, fill_value=None):
if isinstance(x, KerasTensor):
shape = list(x.shape)
if fill_value:
for i, e in enumerate(shape):
if e is None:
shape[i] = fill_value
return np.empty(
shape=shape,
dtype=x.dtype,
)
return x

args_1, kwargs_1 = nest.map_structure(
lambda x: convert_keras_tensor_to_numpy(x, fill_value=83),
(args, kwargs),
)
outputs_1 = fn(*args_1, **kwargs_1)

outputs = outputs_1

if none_in_shape:
args_2, kwargs_2 = nest.map_structure(
lambda x: convert_keras_tensor_to_numpy(x, fill_value=89),
(args, kwargs),
)
outputs_2 = fn(*args_2, **kwargs_2)

flat_out_1 = nest.flatten(outputs_1)
flat_out_2 = nest.flatten(outputs_2)

flat_out = []
for x1, x2 in zip(flat_out_1, flat_out_2):
shape = list(x1.shape)
for i, e in enumerate(x2.shape):
if e != shape[i]:
shape[i] = None
flat_out.append(KerasTensor(shape, standardize_dtype(x1.dtype)))
outputs = nest.pack_sequence_as(outputs_1, flat_out)

def convert_numpy_to_keras_tensor(x):
if is_tensor(x):
return KerasTensor(x.shape, standardize_dtype(x.dtype))
return x

output_spec = nest.map_structure(convert_numpy_to_keras_tensor, outputs)
return output_spec


def scatter(indices, values, shape):
indices = convert_to_tensor(indices)
values = convert_to_tensor(values)
zeros = np.zeros(shape, dtype=values.dtype)

index_length = indices.shape[-1]
value_shape = shape[index_length:]
indices = np.reshape(indices, [-1, index_length])
values = np.reshape(values, [-1] + list(value_shape))

for i in range(indices.shape[0]):
index = indices[i]
zeros[tuple(index)] += values[i]
return zeros


def scatter_update(inputs, indices, updates):
indices = np.array(indices)
indices = np.transpose(indices)
inputs[tuple(indices)] = updates
return inputs


def slice(inputs, start_indices, lengths):
# Validate inputs
assert len(start_indices) == len(lengths)

# Generate list of indices arrays for each dimension
indices = [
np.arange(start, start + length)
for start, length in zip(start_indices, lengths)
]

# Use np.ix_ to create a multidimensional index array
mesh = np.ix_(*indices)

return inputs[mesh]


def slice_update(inputs, start_indices, updates):
# Generate list of indices arrays for each dimension
indices = [
np.arange(start, start + length)
for start, length in zip(start_indices, updates.shape)
]

# Use np.ix_ to create a multidimensional index array
mesh = np.ix_(*indices)
inputs[mesh] = updates
return inputs


def while_loop(
cond,
body,
loop_vars,
maximum_iterations=None,
):
current_iter = 0
iteration_check = (
lambda iter: maximum_iterations is None or iter < maximum_iterations
)
loop_vars = tuple([convert_to_tensor(v) for v in loop_vars])
while cond(*loop_vars) and iteration_check(current_iter):
loop_vars = body(*loop_vars)
if not isinstance(loop_vars, (list, tuple)):
loop_vars = (loop_vars,)
loop_vars = tuple(loop_vars)
current_iter += 1
return loop_vars


def stop_gradient():
raise NotADirectoryError("stop_gradient is not implemented for NumPy.")
45 changes: 45 additions & 0 deletions keras_core/backend/numpy/image.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,45 @@
import jax
import numpy as np

RESIZE_METHODS = (
"bilinear",
"nearest",
"lanczos3",
"lanczos5",
"bicubic",
)


def resize(
image, size, method="bilinear", antialias=False, data_format="channels_last"
):
if method not in RESIZE_METHODS:
raise ValueError(
"Invalid value for argument `method`. Expected of one "
f"{RESIZE_METHODS}. Received: method={method}"
)
if not len(size) == 2:
raise ValueError(
"Argument `size` must be a tuple of two elements "
f"(height, width). Received: size={size}"
)
size = tuple(size)
if len(image.shape) == 4:
if data_format == "channels_last":
size = (image.shape[0],) + size + (image.shape[-1],)
else:
size = (image.shape[0], image.shape[1]) + size
elif len(image.shape) == 3:
if data_format == "channels_last":
size = size + (image.shape[-1],)
else:
size = (image.shape[0],) + size
else:
raise ValueError(
"Invalid input rank: expected rank 3 (single image) "
"or rank 4 (batch of images). Received input with shape: "
f"image.shape={image.shape}"
)
return np.array(
jax.image.resize(image, size, method=method, antialias=antialias)
)
3 changes: 3 additions & 0 deletions keras_core/backend/numpy/layer.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,3 @@
class NumpyLayer:
ariG23498 marked this conversation as resolved.
Show resolved Hide resolved
def _post_build(self):
pass
Loading
Loading