Skip to content

Commit a36a88d

Browse files
authored
Merge 989f655 into 44c1ed3
2 parents 44c1ed3 + 989f655 commit a36a88d

File tree

10 files changed

+281
-8
lines changed

10 files changed

+281
-8
lines changed

doc/reference/functional.rst

Lines changed: 14 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,14 @@
1+
Functional programming
2+
======================
3+
4+
.. https://numpy.org/doc/stable/reference/routines.functional.html
5+
6+
.. autosummary::
7+
:toctree: generated/
8+
:nosignatures:
9+
10+
dpnp.apply_along_axis
11+
dpnp.apply_over_axes
12+
dpnp.vectorize
13+
dpnp.frompyfunc
14+
dpnp.piecewise

doc/reference/manipulation.rst

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -1,4 +1,4 @@
1-
Array Manipulation Routines
1+
Array manipulation routines
22
===========================
33

44
.. https://numpy.org/doc/stable/reference/routines.array-manipulation.html

doc/reference/routines.rst

Lines changed: 3 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -4,7 +4,7 @@ Routines
44

55
The following pages describe NumPy-compatible routines.
66
These functions cover a subset of
7-
`NumPy routines <https://docs.scipy.org/doc/numpy/reference/routines.html>`_.
7+
`NumPy routines <https://numpy.org/doc/stable/reference/routines.html>`_.
88

99
.. currentmodule:: dpnp
1010

@@ -13,10 +13,11 @@ These functions cover a subset of
1313

1414
creation
1515
manipulation
16-
indexing
1716
binary
1817
dtype
1918
fft
19+
functional
20+
indexing
2021
linalg
2122
logic
2223
math

dpnp/dpnp_iface.py

Lines changed: 3 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -81,6 +81,8 @@
8181
from dpnp.dpnp_iface_bitwise import __all__ as __all__bitwise
8282
from dpnp.dpnp_iface_counting import *
8383
from dpnp.dpnp_iface_counting import __all__ as __all__counting
84+
from dpnp.dpnp_iface_functional import *
85+
from dpnp.dpnp_iface_functional import __all__ as __all__functional
8486
from dpnp.dpnp_iface_histograms import *
8587
from dpnp.dpnp_iface_histograms import __all__ as __all__histograms
8688
from dpnp.dpnp_iface_indexing import *
@@ -116,6 +118,7 @@
116118
__all__ += __all__arraycreation
117119
__all__ += __all__bitwise
118120
__all__ += __all__counting
121+
__all__ += __all__functional
119122
__all__ += __all__histograms
120123
__all__ += __all__indexing
121124
__all__ += __all__libmath

dpnp/dpnp_iface_functional.py

Lines changed: 185 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,185 @@
1+
# *****************************************************************************
2+
# Copyright (c) 2024, Intel Corporation
3+
# All rights reserved.
4+
#
5+
# Redistribution and use in source and binary forms, with or without
6+
# modification, are permitted provided that the following conditions are met:
7+
# - Redistributions of source code must retain the above copyright notice,
8+
# this list of conditions and the following disclaimer.
9+
# - Redistributions in binary form must reproduce the above copyright notice,
10+
# this list of conditions and the following disclaimer in the documentation
11+
# and/or other materials provided with the distribution.
12+
#
13+
# THIS SOFTWARE IS PROVIDED BY THE COPYRIGHT HOLDERS AND CONTRIBUTORS "AS IS"
14+
# AND ANY EXPRESS OR IMPLIED WARRANTIES, INCLUDING, BUT NOT LIMITED TO, THE
15+
# IMPLIED WARRANTIES OF MERCHANTABILITY AND FITNESS FOR A PARTICULAR PURPOSE
16+
# ARE DISCLAIMED. IN NO EVENT SHALL THE COPYRIGHT HOLDER OR CONTRIBUTORS BE
17+
# LIABLE FOR ANY DIRECT, INDIRECT, INCIDENTAL, SPECIAL, EXEMPLARY, OR
18+
# CONSEQUENTIAL DAMAGES (INCLUDING, BUT NOT LIMITED TO, PROCUREMENT OF
19+
# SUBSTITUTE GOODS OR SERVICES; LOSS OF USE, DATA, OR PROFITS; OR BUSINESS
20+
# INTERRUPTION) HOWEVER CAUSED AND ON ANY THEORY OF LIABILITY, WHETHER IN
21+
# CONTRACT, STRICT LIABILITY, OR TORT (INCLUDING NEGLIGENCE OR OTHERWISE)
22+
# ARISING IN ANY WAY OUT OF THE USE OF THIS SOFTWARE, EVEN IF ADVISED OF
23+
# THE POSSIBILITY OF SUCH DAMAGE.
24+
# *****************************************************************************
25+
26+
"""
27+
Interface of the functional programming routines part of the DPNP
28+
29+
Notes
30+
-----
31+
This module is a face or public interface file for the library
32+
it contains:
33+
- Interface functions
34+
- documentation for the functions
35+
- The functions parameters check
36+
37+
"""
38+
39+
40+
import numpy
41+
from dpctl.tensor._numpy_helper import normalize_axis_index
42+
43+
import dpnp
44+
45+
__all__ = ["apply_along_axis"]
46+
47+
48+
def apply_along_axis(func1d, axis, arr, *args, **kwargs):
49+
"""
50+
Apply a function to 1-D slices along the given axis.
51+
52+
Execute ``func1d(a, *args, **kwargs)`` where `func1d` operates on
53+
1-D arrays and `a` is a 1-D slice of `arr` along `axis`.
54+
55+
This is equivalent to (but faster than) the following use of
56+
:obj:`dpnp.ndindex` and :obj:`dpnp.s_`, which sets each of
57+
``ii``, ``jj``, and ``kk`` to a tuple of indices::
58+
59+
Ni, Nk = a.shape[:axis], a.shape[axis+1:]
60+
for ii in ndindex(Ni):
61+
for kk in ndindex(Nk):
62+
f = func1d(arr[ii + s_[:,] + kk])
63+
Nj = f.shape
64+
for jj in ndindex(Nj):
65+
out[ii + jj + kk] = f[jj]
66+
67+
Equivalently, eliminating the inner loop, this can be expressed as::
68+
69+
Ni, Nk = a.shape[:axis], a.shape[axis+1:]
70+
for ii in ndindex(Ni):
71+
for kk in ndindex(Nk):
72+
out[ii + s_[...,] + kk] = func1d(arr[ii + s_[:,] + kk])
73+
74+
Parameters
75+
----------
76+
func1d : function (M,) -> (Nj...)
77+
This function should accept 1-D arrays. It is applied to 1-D
78+
slices of `arr` along the specified axis.
79+
axis : int
80+
Axis along which `arr` is sliced.
81+
arr : {dpnp.ndarray, usm_ndarray} (Ni..., M, Nk...)
82+
Input array.
83+
args : any
84+
Additional arguments to `func1d`.
85+
kwargs : any
86+
Additional named arguments to `func1d`.
87+
88+
Returns
89+
-------
90+
out : dpnp.ndarray (Ni..., Nj..., Nk...)
91+
The output array. The shape of `out` is identical to the shape of
92+
`arr`, except along the `axis` dimension. This axis is removed, and
93+
replaced with new dimensions equal to the shape of the return value
94+
of `func1d`.
95+
96+
See Also
97+
--------
98+
:obj:`dpnp.apply_over_axes` : Apply a function repeatedly over
99+
multiple axes.
100+
101+
Examples
102+
--------
103+
>>> import dpnp as np
104+
>>> def my_func(a): # Average first and last element of a 1-D array
105+
... return (a[0] + a[-1]) * 0.5
106+
>>> b = np.array([[1, 2, 3], [4, 5, 6], [7, 8, 9]])
107+
>>> np.apply_along_axis(my_func, 0, b)
108+
array([4., 5., 6.])
109+
>>> np.apply_along_axis(my_func, 1, b)
110+
array([2., 5., 8.])
111+
112+
For a function that returns a 1D array, the number of dimensions in
113+
`outarr` is the same as `arr`.
114+
115+
>>> b = np.array([[8, 1, 7], [4, 3, 9], [5, 2, 6]])
116+
>>> np.apply_along_axis(sorted, 1, b)
117+
array([[1, 7, 8],
118+
[3, 4, 9],
119+
[2, 5, 6]])
120+
121+
For a function that returns a higher dimensional array, those dimensions
122+
are inserted in place of the `axis` dimension.
123+
124+
>>> b = np.array([[1, 2, 3], [4, 5, 6], [7, 8, 9]])
125+
>>> np.apply_along_axis(np.diag, -1, b)
126+
array([[[1, 0, 0],
127+
[0, 2, 0],
128+
[0, 0, 3]],
129+
[[4, 0, 0],
130+
[0, 5, 0],
131+
[0, 0, 6]],
132+
[[7, 0, 0],
133+
[0, 8, 0],
134+
[0, 0, 9]]])
135+
136+
"""
137+
138+
dpnp.check_supported_arrays_type(arr)
139+
nd = arr.ndim
140+
exec_q = arr.sycl_queue
141+
usm_type = arr.usm_type
142+
axis = normalize_axis_index(axis, nd)
143+
144+
# arr, with the iteration axis at the end
145+
inarr_view = dpnp.moveaxis(arr, axis, -1)
146+
147+
# compute indices for the iteration axes, and append a trailing ellipsis to
148+
# prevent 0d arrays decaying to scalars
149+
# TODO: replace with dpnp.ndindex
150+
inds = numpy.ndindex(inarr_view.shape[:-1])
151+
inds = (ind + (Ellipsis,) for ind in inds)
152+
153+
# invoke the function on the first item
154+
try:
155+
ind0 = next(inds)
156+
except StopIteration:
157+
raise ValueError(
158+
"Cannot apply_along_axis when any iteration dimensions are 0"
159+
) from None
160+
res = dpnp.asanyarray(
161+
func1d(inarr_view[ind0], *args, **kwargs),
162+
sycl_queue=exec_q,
163+
usm_type=usm_type,
164+
)
165+
166+
# build a buffer for storing evaluations of func1d.
167+
# remove the requested axis, and add the new ones on the end.
168+
# laid out so that each write is contiguous.
169+
# for a tuple index inds, buff[inds] = func1d(inarr_view[inds])
170+
buff = dpnp.empty_like(res, shape=inarr_view.shape[:-1] + res.shape)
171+
172+
# save the first result, then compute and save all remaining results
173+
buff[ind0] = res
174+
for ind in inds:
175+
buff[ind] = dpnp.asanyarray(
176+
func1d(inarr_view[ind], *args, **kwargs),
177+
sycl_queue=exec_q,
178+
usm_type=usm_type,
179+
)
180+
181+
# restore the inserted axes back to where they belong
182+
for _ in range(res.ndim):
183+
buff = dpnp.moveaxis(buff, -1, axis)
184+
185+
return buff

dpnp/dpnp_iface_manipulation.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -25,7 +25,7 @@
2525
# *****************************************************************************
2626

2727
"""
28-
Interface of the Array manipulation routines part of the DPNP
28+
Interface of the array manipulation routines part of the DPNP
2929
3030
Notes
3131
-----

tests/test_functional.py

Lines changed: 51 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,51 @@
1+
import numpy
2+
import pytest
3+
from numpy.testing import assert_array_equal, assert_equal, assert_raises
4+
5+
import dpnp
6+
7+
from .helper import get_all_dtypes
8+
9+
10+
class TestApplyAlongAxis:
11+
def test_tuple_func1d(self):
12+
def sample_1d(x):
13+
return x[1], x[0]
14+
15+
a = numpy.array([[1, 2], [3, 4]])
16+
ia = dpnp.array(a)
17+
18+
# 2d insertion along first axis
19+
expected = numpy.apply_along_axis(sample_1d, 1, a)
20+
result = dpnp.apply_along_axis(sample_1d, 1, ia)
21+
assert_array_equal(result, expected)
22+
23+
@pytest.mark.parametrize("stride", [-1, 2, -3])
24+
def test_stride(self, stride):
25+
a = numpy.ones((20, 10), dtype="f")
26+
ia = dpnp.array(a)
27+
28+
expected = numpy.apply_along_axis(len, 0, a[::stride, ::stride])
29+
result = dpnp.apply_along_axis(len, 0, ia[::stride, ::stride])
30+
assert_array_equal(result, expected)
31+
32+
@pytest.mark.parametrize("dtype", get_all_dtypes())
33+
def test_args(self, dtype):
34+
a = numpy.ones((20, 10))
35+
ia = dpnp.array(a)
36+
37+
# kwargs
38+
expected = numpy.apply_along_axis(
39+
numpy.mean, 0, a, dtype=dtype, keepdims=True
40+
)
41+
result = dpnp.apply_along_axis(
42+
dpnp.mean, 0, ia, dtype=dtype, keepdims=True
43+
)
44+
assert_array_equal(result, expected)
45+
46+
# positional args: axis, dtype, out, keepdims
47+
expected = numpy.apply_along_axis(
48+
numpy.mean, 0, a, 0, dtype, None, True
49+
)
50+
result = dpnp.apply_along_axis(dpnp.mean, 0, ia, 0, dtype, None, True)
51+
assert_array_equal(result, expected)

tests/test_sycl_queue.py

Lines changed: 12 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -2172,6 +2172,18 @@ def test_split(func, data1, device):
21722172
assert_sycl_queue_equal(result[1].sycl_queue, x1.sycl_queue)
21732173

21742174

2175+
@pytest.mark.parametrize(
2176+
"device",
2177+
valid_devices,
2178+
ids=[device.filter_string for device in valid_devices],
2179+
)
2180+
def test_apply_along_axis(device):
2181+
x = dpnp.arange(9, device=device).reshape(3, 3)
2182+
result = dpnp.apply_along_axis(dpnp.sum, 0, x)
2183+
2184+
assert_sycl_queue_equal(result.sycl_queue, x.sycl_queue)
2185+
2186+
21752187
@pytest.mark.parametrize(
21762188
"device_x",
21772189
valid_devices,

tests/test_usm_type.py

Lines changed: 8 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -767,6 +767,14 @@ def test_2in_with_scalar_1out(func, data, scalar, usm_type):
767767
assert z.usm_type == usm_type
768768

769769

770+
@pytest.mark.parametrize("usm_type", list_of_usm_types, ids=list_of_usm_types)
771+
def test_apply_along_axis(usm_type):
772+
x = dp.arange(9, usm_type=usm_type).reshape(3, 3)
773+
y = dp.apply_along_axis(dp.sum, 0, x)
774+
775+
assert x.usm_type == y.usm_type
776+
777+
770778
@pytest.mark.parametrize("usm_type", list_of_usm_types, ids=list_of_usm_types)
771779
def test_broadcast_to(usm_type):
772780
x = dp.ones(7, usm_type=usm_type)

tests/third_party/cupy/lib_tests/test_shape_base.py

Lines changed: 3 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -5,15 +5,15 @@
55
from dpctl.tensor._numpy_helper import AxisError
66

77
import dpnp as cupy
8+
from tests.helper import has_support_aspect64
89
from tests.third_party.cupy import testing
910

1011

1112
@testing.parameterize(*(testing.product({"axis": [0, 1, -1]})))
12-
@pytest.mark.skip("'apply_along_axis' is not implemented yet")
1313
class TestApplyAlongAxis(unittest.TestCase):
1414
@testing.numpy_cupy_array_equal()
1515
def test_simple(self, xp):
16-
a = xp.ones((20, 10), "d")
16+
a = xp.ones((20, 10), dtype="f")
1717
return xp.apply_along_axis(len, self.axis, a)
1818

1919
@testing.for_all_dtypes(no_bool=True)
@@ -22,7 +22,7 @@ def test_3d(self, xp, dtype):
2222
a = xp.arange(27, dtype=dtype).reshape((3, 3, 3))
2323
return xp.apply_along_axis(xp.sum, self.axis, a)
2424

25-
@testing.numpy_cupy_array_equal()
25+
@testing.numpy_cupy_array_equal(type_check=has_support_aspect64())
2626
def test_0d_array(self, xp):
2727

2828
def sum_to_0d(x):
@@ -100,7 +100,6 @@ def func(x):
100100

101101

102102
@testing.with_requires("numpy>=1.16")
103-
@pytest.mark.skip("'apply_along_axis' is not implemented yet")
104103
def test_apply_along_axis_invalid_axis():
105104
for xp in [numpy, cupy]:
106105
a = xp.ones((8, 4))

0 commit comments

Comments
 (0)