Skip to content

Commit 8a2ffee

Browse files
authored
[Unity] Do not import SciPy by default (#16136)
SciPy recently got a few conflicts with numpy and emits warning messages at loading time.
1 parent 44d80e5 commit 8a2ffee

File tree

9 files changed

+52
-27
lines changed

9 files changed

+52
-27
lines changed

python/tvm/autotvm/tuner/__init__.py

Lines changed: 2 additions & 5 deletions
Original file line numberDiff line numberDiff line change
@@ -22,10 +22,7 @@
2222
"""
2323

2424
from . import callback
25-
26-
from .tuner import Tuner
27-
28-
from .index_based_tuner import GridSearchTuner, RandomTuner
2925
from .ga_tuner import GATuner
26+
from .index_based_tuner import GridSearchTuner, RandomTuner
27+
from .tuner import Tuner
3028
from .xgboost_tuner import XGBTuner
31-
from .droplet_turner import DropletTuner

python/tvm/relax/dpl/pattern.py

Lines changed: 2 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -26,7 +26,7 @@
2626
import tvm._ffi as tvm_ffi
2727
from tvm.ir.container import Array
2828
from tvm.ir.expr import PrimExpr
29-
from tvm.relay.op import get
29+
from tvm.ir.op import Op
3030

3131
from ...ir import make_node
3232
from ...ir.base import Node
@@ -756,7 +756,7 @@ def is_op(op_name: str) -> ExprPattern:
756756
result: tvm.relax.dpl.ExprPattern
757757
The resulting ExprPattern
758758
"""
759-
op = get(op_name)
759+
op = Op.get(op_name)
760760
return ExprPattern(op)
761761

762762

python/tvm/relay/analysis/sparse_conv2d.py

Lines changed: 7 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -21,11 +21,12 @@
2121
to block sparse model
2222
"""
2323
from collections import namedtuple
24+
2425
import numpy as np
25-
import scipy.sparse as sp
26+
2627
import tvm
27-
from . import _ffi_api
2828

29+
from . import _ffi_api
2930

3031
SparseAnalysisResult = namedtuple(
3132
"SparseAnalysisResult",
@@ -79,9 +80,11 @@ def process_params(
7980
"""
8081

8182
# pylint: disable=import-outside-toplevel
82-
from tvm.auto_scheduler.search_task import (
83+
import scipy.sparse as sp
84+
85+
from tvm.auto_scheduler.search_task import ( # lazily import to avoid recursive dependency
8386
register_task_input_buffer,
84-
) # lazily import to avoid recursive dependency
87+
)
8588

8689
memo = SparseAnalysisResult(weight_name=[], weight_shape=[])
8790
weight_names = _search_conv2d_op_weight(expr)

python/tvm/relay/analysis/sparse_dense.py

Lines changed: 7 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -21,11 +21,12 @@
2121
to block sparse model
2222
"""
2323
from collections import namedtuple
24+
2425
import numpy as np
25-
import scipy.sparse as sp
26+
2627
import tvm
27-
from . import _ffi_api
2828

29+
from . import _ffi_api
2930

3031
SparseAnalysisResult = namedtuple(
3132
"SparseAnalysisResult",
@@ -75,9 +76,11 @@ def process_params(expr, params, block_size, sparsity_threshold):
7576
"""
7677

7778
# pylint: disable=import-outside-toplevel
78-
from tvm.auto_scheduler.search_task import (
79+
import scipy.sparse as sp
80+
81+
from tvm.auto_scheduler.search_task import ( # lazily import to avoid recursive dependency
7982
register_task_input_buffer,
80-
) # lazily import to avoid recursive dependency
83+
)
8184

8285
memo = SparseAnalysisResult(weight_name=[], weight_shape=[])
8386
weight_names = _search_dense_op_weight(expr)

python/tvm/relay/frontend/tensorflow_ops.py

Lines changed: 3 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -1481,10 +1481,10 @@ def _impl(inputs, attr, params, mod):
14811481

14821482

14831483
def _sparse_tensor_dense_add():
1484-
# Sparse utility from scipy
1485-
from scipy.sparse import csr_matrix
1486-
14871484
def _impl(inputs, attr, params, mod):
1485+
# Sparse utility from scipy
1486+
from scipy.sparse import csr_matrix
1487+
14881488
assert (
14891489
len(inputs) == 4
14901490
), "There should be 4 input tensors [sparse_indices, sparse_values, sparse_shape, dense]."

python/tvm/relay/qnn/op/legalizations.py

Lines changed: 17 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -17,7 +17,7 @@
1717
# pylint: disable=invalid-name, unused-argument
1818
"""Backend QNN related feature registration"""
1919
import numpy as np
20-
from scipy import special
20+
2121
import tvm
2222
from tvm import relay
2323
from tvm._ffi.base import TVMError
@@ -78,14 +78,29 @@ def hardswish_func(x):
7878
register_qnn_unary_op_legalize("qnn.sqrt", np.sqrt)
7979
register_qnn_unary_op_legalize("qnn.rsqrt", lambda arr: 1 / np.sqrt(arr))
8080
register_qnn_unary_op_legalize("qnn.exp", np.exp)
81-
register_qnn_unary_op_legalize("qnn.erf", special.erf)
8281
register_qnn_unary_op_legalize("qnn.sigmoid", lambda arr: 1 / (1 + np.exp(-arr)))
8382
register_qnn_unary_op_legalize("qnn.hardswish", hardswish_func)
8483
register_qnn_unary_op_legalize("qnn.tanh", np.tanh)
8584
register_qnn_unary_op_legalize("qnn.log", np.log)
8685
register_qnn_unary_op_legalize("qnn.abs", np.abs)
8786

8887

88+
@reg.register_qnn_legalize("qnn.erf")
89+
def _legalize_qnn_erf(attrs, inputs, types):
90+
from scipy import special # pylint: disable=import-outside-toplevel
91+
92+
return create_integer_lookup_op(
93+
input_arg=inputs[0],
94+
floating_point_func=special.erf,
95+
in_scale=inputs[1],
96+
in_zero_point=inputs[2],
97+
out_scale=inputs[3],
98+
out_zero_point=inputs[4],
99+
in_dtype=types[0].dtype,
100+
out_dtype=types[0].dtype,
101+
)
102+
103+
89104
# Default to None. If overridden by target, this will not be run.
90105
# Generic QNN Conv2D legalization function.
91106
@tvm.target.generic_func

python/tvm/topi/arm_cpu/qnn_legalize.py

Lines changed: 5 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -30,11 +30,12 @@
3030
"""
3131

3232
import numpy as np
33-
from scipy.signal import convolve2d
34-
from tvm.topi.utils import get_const_tuple
33+
3534
from tvm import nd, relay
36-
from .qnn_alter_op import prev_ops_match, edit_attrs
35+
from tvm.topi.utils import get_const_tuple
36+
3737
from ..nn import bias_add_legalize
38+
from .qnn_alter_op import edit_attrs, prev_ops_match
3839

3940

4041
def _compute_fixed_conv2d_outputs(requantize_op):
@@ -112,6 +113,7 @@ def _compute_fixed_depthwise_outputs(requantize_op, fixed_channel_inputs):
112113
an output channel index, and each value is the value that all entries in that output channel
113114
will have. If the block has no fixed output channels, this dictionary will be empty.
114115
"""
116+
from scipy.signal import convolve2d # pylint: disable=import-outside-toplevel
115117

116118
bias_add_op = requantize_op.args[0]
117119
depthwise_op = bias_add_op.args[0]

python/tvm/topi/cuda/sparse.py

Lines changed: 5 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -17,13 +17,12 @@
1717

1818
"""Sparse operators"""
1919
import numpy as np
20-
import scipy.sparse as sp
2120

2221
import tvm
2322
from tvm import relay, te
2423

2524
from .. import nn
26-
from ..utils import traverse_inline, get_const_tuple, prod, get_const_int, ceil_div
25+
from ..utils import ceil_div, get_const_int, get_const_tuple, prod, traverse_inline
2726
from .transform import schedule_transpose_from_existing
2827

2928

@@ -358,6 +357,8 @@ def schedule_sparse_dense_padded(outs):
358357

359358
def pad_sparse_matrix(matrix, blocksize):
360359
"""Pad rows of sparse matrix matrix so that they are a multiple of blocksize."""
360+
import scipy.sparse as sp # pylint: disable=import-outside-toplevel
361+
361362
assert isinstance(matrix, sp.bsr_matrix)
362363
new_entries = np.zeros(matrix.shape[0], dtype=matrix.indptr.dtype)
363364
bsr = matrix.blocksize[0]
@@ -394,6 +395,8 @@ def _alter_sparse_dense_layout(_attrs, inputs, _tinfos, _out_type):
394395
sparse_dense implementation for one that operates on a padded matrix. We
395396
also pad the matrix.
396397
"""
398+
import scipy.sparse as sp # pylint: disable=import-outside-toplevel
399+
397400
# TODO(ANSHUMAN87): Handle for sparse_lhs case too
398401
if (
399402
isinstance(inputs[1], relay.Constant)

python/tvm/topi/nn/conv3d_transpose.py

Lines changed: 4 additions & 2 deletions
Original file line numberDiff line numberDiff line change
@@ -18,11 +18,11 @@
1818
"""Transposed 3D convolution operators (sometimes called Deconvolution)."""
1919
import tvm
2020
from tvm import te
21-
from tvm import relay
21+
22+
from ..utils import simplify
2223
from .dilate import dilate
2324
from .pad import pad
2425
from .utils import get_pad_tuple3d
25-
from ..utils import simplify
2626

2727

2828
def conv3d_transpose_ncdhw(Input, Filter, strides, padding, out_dtype, output_padding):
@@ -143,6 +143,8 @@ def conv3d_transpose_legalize(attrs, inputs, types):
143143
result : tvm.relay.Expr
144144
The legalized expr
145145
"""
146+
from tvm import relay # pylint: disable=import-outside-toplevel
147+
146148
if attrs["data_layout"] == "NDHWC":
147149
data, kernel = inputs
148150
kernel_layout = attrs["kernel_layout"]

0 commit comments

Comments
 (0)