Skip to content
GitLab
Explore
Sign in
Primary navigation
Search or go to…
Project
T
tvm
Manage
Activity
Members
Labels
Plan
Issues
Issue boards
Milestones
Wiki
Code
Merge requests
Repository
Branches
Commits
Tags
Repository graph
Compare revisions
Snippets
Build
Pipelines
Jobs
Pipeline schedules
Artifacts
Deploy
Releases
Model registry
Operate
Environments
Monitor
Incidents
Service Desk
Analyze
Value stream analytics
Contributor analytics
CI/CD analytics
Repository analytics
Model experiments
Help
Help
Support
GitLab documentation
Compare GitLab plans
Community forum
Contribute to GitLab
Provide feedback
Terms and privacy
Keyboard shortcuts
?
Snippets
Groups
Projects
Show more breadcrumbs
cld
ml
tvm
Commits
833855e7
Commit
833855e7
authored
7 years ago
by
Tianqi Chen
Committed by
GitHub
7 years ago
Browse files
Options
Downloads
Patches
Plain Diff
[TOPI] Fix reduction fusion with injective input (#475)
parent
203b8188
No related branches found
No related tags found
No related merge requests found
Changes
2
Hide whitespace changes
Inline
Side-by-side
Showing
2 changed files
topi/python/topi/cuda/reduction.py
+13
-7
13 additions, 7 deletions
topi/python/topi/cuda/reduction.py
topi/tests/python/test_topi_reduce.py
+9
-8
9 additions, 8 deletions
topi/tests/python/test_topi_reduce.py
with
22 additions
and
15 deletions
topi/python/topi/cuda/reduction.py
+
13
−
7
View file @
833855e7
...
...
@@ -34,6 +34,7 @@ def _schedule_reduce(op, sch):
# Bind the axes to threads and blocks
sch
[
data_out
].
bind
(
sch
[
data_out
].
op
.
reduce_axis
[
0
],
thread_x
)
sch
[
data_out
].
set_store_predicate
(
thread_x
.
equal
(
0
))
sch
[
data_out
].
bind
(
outer_in
,
thread_y
)
sch
[
data_out
].
bind
(
bx
,
block_x
)
else
:
...
...
@@ -57,17 +58,22 @@ def schedule_reduce(outs):
"""
outs
=
[
outs
]
if
isinstance
(
outs
,
tvm
.
tensor
.
Tensor
)
else
outs
sch
=
tvm
.
create_schedule
([
x
.
op
for
x
in
outs
])
def
traverse
(
operator
):
def
traverse_before_reduce
(
operator
):
if
tag
.
is_injective
(
operator
.
tag
):
if
operator
not
in
sch
.
outputs
:
sch
[
operator
].
compute_inline
()
for
tensor
in
operator
.
input_tensors
:
if
tensor
.
op
.
input_tensors
:
traverse
(
tensor
.
op
)
sch
[
operator
].
compute_inline
()
else
:
raise
RuntimeError
(
"
Unsupported operator: %s
"
%
operator
.
tag
)
def
traverse_after_reduce
(
operator
):
if
tag
.
is_broadcast
(
operator
.
tag
):
raise
RuntimeError
(
"
Not yet support ewise after reduce
"
)
elif
operator
.
tag
==
'
comm_reduce
'
:
_schedule_reduce
(
operator
,
sch
)
for
tensor
in
operator
.
input_tensors
:
traverse_before_reduce
(
tensor
.
op
)
else
:
raise
RuntimeError
(
"
Unsupported operator: %s
"
%
operator
.
tag
)
traverse
(
outs
[
0
].
op
)
traverse
_after_reduce
(
outs
[
0
].
op
)
return
sch
This diff is collapsed.
Click to expand it.
topi/tests/python/test_topi_reduce.py
+
9
−
8
View file @
833855e7
...
...
@@ -7,12 +7,13 @@ import topi
def
verify_reduce_map_ele
(
in_shape
,
axis
,
keepdims
,
type
=
"
sum
"
):
# Build the logic and compile the function
A
=
tvm
.
placeholder
(
shape
=
in_shape
,
name
=
"
A
"
)
A1
=
topi
.
exp
(
A
)
if
type
==
"
sum
"
:
B
=
topi
.
sum
(
A
,
axis
=
axis
,
keepdims
=
keepdims
)
B
=
topi
.
sum
(
A
1
,
axis
=
axis
,
keepdims
=
keepdims
)
elif
type
==
"
max
"
:
B
=
topi
.
max
(
A
,
axis
=
axis
,
keepdims
=
keepdims
)
B
=
topi
.
max
(
A
1
,
axis
=
axis
,
keepdims
=
keepdims
)
elif
type
==
"
min
"
:
B
=
topi
.
min
(
A
,
axis
=
axis
,
keepdims
=
keepdims
)
B
=
topi
.
min
(
A
1
,
axis
=
axis
,
keepdims
=
keepdims
)
else
:
raise
NotImplementedError
s
=
topi
.
cuda
.
schedule_reduce
(
B
)
...
...
@@ -23,15 +24,15 @@ def verify_reduce_map_ele(in_shape, axis, keepdims, type="sum"):
return
ctx
=
tvm
.
gpu
(
0
)
if
device
==
"
cuda
"
else
tvm
.
cl
(
0
)
foo
=
tvm
.
build
(
s
,
[
A
,
B
],
device
,
name
=
"
sum
"
)
# Test
in_npy
=
np
.
random
.
normal
(
size
=
in_shape
).
astype
(
np
.
float32
)
in_npy
=
np
.
random
.
uniform
(
size
=
in_shape
).
astype
(
np
.
float32
)
in_npy_map
=
np
.
exp
(
in_npy
)
if
type
==
"
sum
"
:
out_npy
=
in_npy
.
sum
(
axis
=
axis
,
keepdims
=
keepdims
)
out_npy
=
in_npy
_map
.
sum
(
axis
=
axis
,
keepdims
=
keepdims
)
elif
type
==
"
max
"
:
out_npy
=
in_npy
.
max
(
axis
=
axis
,
keepdims
=
keepdims
)
out_npy
=
in_npy
_map
.
max
(
axis
=
axis
,
keepdims
=
keepdims
)
elif
type
==
"
min
"
:
out_npy
=
in_npy
.
min
(
axis
=
axis
,
keepdims
=
keepdims
)
out_npy
=
in_npy
_map
.
min
(
axis
=
axis
,
keepdims
=
keepdims
)
else
:
raise
NotImplementedError
...
...
This diff is collapsed.
Click to expand it.
Preview
0%
Loading
Try again
or
attach a new file
.
Cancel
You are about to add
0
people
to the discussion. Proceed with caution.
Finish editing this message first!
Save comment
Cancel
Please
register
or
sign in
to comment