Skip to content

Commit 0a7f42a

Browse files
committed
Fix and update docs for utils.jl
1 parent 898b558 commit 0a7f42a

File tree

5 files changed

+79
-42
lines changed

5 files changed

+79
-42
lines changed

docs/src/utilities.md

Lines changed: 2 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -42,7 +42,9 @@ Flux.orthogonal
4242
Flux.sparse_init
4343
Flux.identity_init
4444
Flux.ones32
45+
Flux.zeros32
4546
Flux.rand32
47+
Flux.randn32
4648
```
4749

4850
## Changing the type of model parameters

src/layers/basic.jl

Lines changed: 3 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -155,7 +155,7 @@ struct Dense{F, M<:AbstractMatrix, B}
155155
bias::B
156156
σ::F
157157
function Dense(W::M, bias = true, σ::F = identity) where {M<:AbstractMatrix, F}
158-
b = create_bias(W, bias, size(W,1))
158+
b = _create_bias(W, bias, size(W,1))
159159
new{F,M,typeof(b)}(W, b, σ)
160160
end
161161
end
@@ -228,7 +228,7 @@ struct Scale{F, A<:AbstractArray, B}
228228
bias::B
229229
σ::F
230230
function Scale(scale::A, bias::B = true, σ::F = identity) where {A<:AbstractArray, B<:Union{Bool, AbstractArray}, F}
231-
b = create_bias(scale, bias, size(scale)...)
231+
b = _create_bias(scale, bias, size(scale)...)
232232
new{F, A, typeof(b)}(scale, b, σ)
233233
end
234234
end
@@ -403,7 +403,7 @@ struct Bilinear{F,A,B}
403403
σ::F
404404
function Bilinear(W::A, bias = true, σ::F = identity) where {A<:AbstractArray, F}
405405
ndims(A) == 3 || throw(ArgumentError("expected a 3-array of weights"))
406-
b = create_bias(W, bias, size(W,1))
406+
b = _create_bias(W, bias, size(W,1))
407407
new{F,A,typeof(b)}(W, b, σ)
408408
end
409409
end

src/layers/conv.jl

Lines changed: 3 additions & 3 deletions
Original file line numberDiff line numberDiff line change
@@ -156,7 +156,7 @@ function Conv(w::AbstractArray{T,N}, b = true, σ = identity;
156156
stride = expand(Val(N-2), stride)
157157
dilation = expand(Val(N-2), dilation)
158158
pad = calc_padding(Conv, pad, size(w)[1:N-2], dilation, stride)
159-
bias = create_bias(w, b, size(w, N))
159+
bias = _create_bias(w, b, size(w, N))
160160
return Conv(σ, w, bias, stride, pad, dilation, groups)
161161
end
162162

@@ -293,7 +293,7 @@ function ConvTranspose(w::AbstractArray{T,N}, bias = true, σ = identity;
293293
stride = expand(Val(N-2), stride)
294294
dilation = expand(Val(N-2), dilation)
295295
pad = calc_padding(ConvTranspose, pad, size(w)[1:N-2], dilation, stride)
296-
b = create_bias(w, bias, size(w, N-1) * groups)
296+
b = _create_bias(w, bias, size(w, N-1) * groups)
297297
return ConvTranspose(σ, w, b, stride, pad, dilation, groups)
298298
end
299299

@@ -441,7 +441,7 @@ function CrossCor(w::AbstractArray{T,N}, bias = true, σ = identity;
441441
stride = expand(Val(N-2), stride)
442442
dilation = expand(Val(N-2), dilation)
443443
pad = calc_padding(CrossCor, pad, size(w)[1:N-2], dilation, stride)
444-
b = create_bias(w, bias, size(w, N))
444+
b = _create_bias(w, bias, size(w, N))
445445
return CrossCor(σ, w, b, stride, pad, dilation)
446446
end
447447

src/layers/normalise.jl

Lines changed: 8 additions & 8 deletions
Original file line numberDiff line numberDiff line change
@@ -10,7 +10,7 @@ _dropout_shape(s, dims) = tuple((i ∉ dims ? 1 : si for (i, si) ∈ enumerate(s
1010
_dropout_kernel(y::T, p, q) where {T} = y > p ? T(1 / q) : T(0)
1111

1212
"""
13-
dropout([rng = rng_from_array(x)], x, p; dims=:, active=true)
13+
dropout([rng = _rng_from_array(x)], x, p; dims=:, active=true)
1414
1515
The dropout function. If `active` is `true`,
1616
for each input, either sets that input to `0` (with probability
@@ -34,7 +34,7 @@ function dropout(rng, x, p; dims=:, active::Bool=true)
3434
y = dropout_mask(rng, x, p, dims=dims)
3535
return x .* y
3636
end
37-
dropout(x, p; kwargs...) = dropout(rng_from_array(x), x, p; kwargs...)
37+
dropout(x, p; kwargs...) = dropout(_rng_from_array(x), x, p; kwargs...)
3838

3939
dropout_mask(rng::CUDA.RNG, x::CuArray, p; kwargs...) = _dropout_mask(rng, x, p; kwargs...)
4040
dropout_mask(rng, x::CuArray, p; kwargs...) =
@@ -51,7 +51,7 @@ end
5151
ChainRulesCore.@non_differentiable dropout_mask(::Any, ::Any, ::Any)
5252

5353
"""
54-
Dropout(p; dims=:, rng = rng_from_array())
54+
Dropout(p; dims=:, rng = _rng_from_array())
5555
5656
Dropout layer.
5757
@@ -96,9 +96,9 @@ mutable struct Dropout{F,D,R<:AbstractRNG}
9696
active::Union{Bool, Nothing}
9797
rng::R
9898
end
99-
Dropout(p, dims, active) = Dropout(p, dims, active, rng_from_array())
99+
Dropout(p, dims, active) = Dropout(p, dims, active, _rng_from_array())
100100

101-
function Dropout(p; dims=:, rng = rng_from_array())
101+
function Dropout(p; dims=:, rng = _rng_from_array())
102102
@assert 0 p 1
103103
Dropout(p, dims, nothing, rng)
104104
end
@@ -121,7 +121,7 @@ function Base.show(io::IO, d::Dropout)
121121
end
122122

123123
"""
124-
AlphaDropout(p; rng = rng_from_array())
124+
AlphaDropout(p; rng = _rng_from_array())
125125
126126
A dropout layer. Used in
127127
[Self-Normalizing Neural Networks](https://arxiv.org/abs/1706.02515).
@@ -155,8 +155,8 @@ mutable struct AlphaDropout{F,R<:AbstractRNG}
155155
new{typeof(p), typeof(rng)}(p, active, rng)
156156
end
157157
end
158-
AlphaDropout(p, active) = AlphaDropout(p, active, rng_from_array())
159-
AlphaDropout(p; rng = rng_from_array()) = AlphaDropout(p, nothing, rng)
158+
AlphaDropout(p, active) = AlphaDropout(p, active, _rng_from_array())
159+
AlphaDropout(p; rng = _rng_from_array()) = AlphaDropout(p, nothing, rng)
160160

161161
@functor AlphaDropout
162162
trainable(a::AlphaDropout) = (;)

src/utils.jl

Lines changed: 63 additions & 28 deletions
Original file line numberDiff line numberDiff line change
@@ -34,7 +34,7 @@ ofeltype(x, y) = convert(float(eltype(x)), y)
3434
epseltype(x) = eps(float(eltype(x)))
3535

3636
"""
37-
rng_from_array([x])
37+
_rng_from_array([x])
3838
3939
Create an instance of the RNG most appropriate for `x`.
4040
The current defaults are:
@@ -43,12 +43,12 @@ The current defaults are:
4343
- Julia version is < 1.7: `Random.GLOBAL_RNG`
4444
- Julia version is >= 1.7: `Random.default_rng()`
4545
"""
46-
rng_from_array(::AbstractArray) = rng_from_array()
47-
rng_from_array(::CuArray) = CUDA.default_rng()
46+
_rng_from_array(::AbstractArray) = _rng_from_array()
47+
_rng_from_array(::CuArray) = CUDA.default_rng()
4848
if VERSION >= v"1.7"
49-
rng_from_array() = Random.default_rng()
49+
_rng_from_array() = Random.default_rng()
5050
else
51-
rng_from_array() = Random.GLOBAL_RNG
51+
_rng_from_array() = Random.GLOBAL_RNG
5252
end
5353

5454
"""
@@ -91,8 +91,8 @@ function glorot_uniform(rng::AbstractRNG, dims::Integer...; gain::Real=1)
9191
scale = Float32(gain) * sqrt(24.0f0 / sum(nfan(dims...)))
9292
(rand(rng, Float32, dims...) .- 0.5f0) .* scale
9393
end
94-
glorot_uniform(dims::Integer...; kw...) = glorot_uniform(rng_from_array(), dims...; kw...)
95-
glorot_uniform(rng::AbstractRNG=rng_from_array(); init_kwargs...) = (dims...; kwargs...) -> glorot_uniform(rng, dims...; init_kwargs..., kwargs...)
94+
glorot_uniform(dims::Integer...; kw...) = glorot_uniform(_rng_from_array(), dims...; kw...)
95+
glorot_uniform(rng::AbstractRNG=_rng_from_array(); init_kwargs...) = (dims...; kwargs...) -> glorot_uniform(rng, dims...; init_kwargs..., kwargs...)
9696

9797
ChainRulesCore.@non_differentiable glorot_uniform(::Any...)
9898

@@ -134,8 +134,8 @@ function glorot_normal(rng::AbstractRNG, dims::Integer...; gain::Real=1)
134134
std = Float32(gain) * sqrt(2.0f0 / sum(nfan(dims...)))
135135
randn(rng, Float32, dims...) .* std
136136
end
137-
glorot_normal(dims::Integer...; kwargs...) = glorot_normal(rng_from_array(), dims...; kwargs...)
138-
glorot_normal(rng::AbstractRNG=rng_from_array(); init_kwargs...) = (dims...; kwargs...) -> glorot_normal(rng, dims...; init_kwargs..., kwargs...)
137+
glorot_normal(dims::Integer...; kwargs...) = glorot_normal(_rng_from_array(), dims...; kwargs...)
138+
glorot_normal(rng::AbstractRNG=_rng_from_array(); init_kwargs...) = (dims...; kwargs...) -> glorot_normal(rng, dims...; init_kwargs..., kwargs...)
139139

140140
ChainRulesCore.@non_differentiable glorot_normal(::Any...)
141141

@@ -169,8 +169,8 @@ function kaiming_uniform(rng::AbstractRNG, dims::Integer...; gain::Real = √2)
169169
return (rand(rng, Float32, dims...) .- 0.5f0) .* 2bound
170170
end
171171

172-
kaiming_uniform(dims::Integer...; kwargs...) = kaiming_uniform(rng_from_array(), dims...; kwargs...)
173-
kaiming_uniform(rng::AbstractRNG=rng_from_array(); init_kwargs...) = (dims...; kwargs...) -> kaiming_uniform(rng, dims...; init_kwargs..., kwargs...)
172+
kaiming_uniform(dims::Integer...; kwargs...) = kaiming_uniform(_rng_from_array(), dims...; kwargs...)
173+
kaiming_uniform(rng::AbstractRNG=_rng_from_array(); init_kwargs...) = (dims...; kwargs...) -> kaiming_uniform(rng, dims...; init_kwargs..., kwargs...)
174174

175175
ChainRulesCore.@non_differentiable kaiming_uniform(::Any...)
176176

@@ -206,7 +206,7 @@ function kaiming_normal(rng::AbstractRNG, dims::Integer...; gain::Real = √2f0)
206206
return randn(rng, Float32, dims...) .* std
207207
end
208208

209-
kaiming_normal(dims::Integer...; kwargs...) = kaiming_normal(rng_from_array(), dims...; kwargs...)
209+
kaiming_normal(dims::Integer...; kwargs...) = kaiming_normal(_rng_from_array(), dims...; kwargs...)
210210
kaiming_normal(rng::AbstractRNG; init_kwargs...) = (dims...; kwargs...) -> kaiming_normal(rng, dims...; init_kwargs..., kwargs...)
211211

212212
ChainRulesCore.@non_differentiable kaiming_normal(::Any...)
@@ -252,8 +252,8 @@ function truncated_normal(rng::AbstractRNG, dims::Integer...; mean = 0, std = 1,
252252
return xs
253253
end
254254

255-
truncated_normal(dims::Integer...; kwargs...) = truncated_normal(rng_from_array(), dims...; kwargs...)
256-
truncated_normal(rng::AbstractRNG=rng_from_array(); init_kwargs...) = (dims...; kwargs...) -> truncated_normal(rng, dims...; init_kwargs..., kwargs...)
255+
truncated_normal(dims::Integer...; kwargs...) = truncated_normal(_rng_from_array(), dims...; kwargs...)
256+
truncated_normal(rng::AbstractRNG=_rng_from_array(); init_kwargs...) = (dims...; kwargs...) -> truncated_normal(rng, dims...; init_kwargs..., kwargs...)
257257

258258
ChainRulesCore.@non_differentiable truncated_normal(::Any...)
259259

@@ -313,8 +313,8 @@ function orthogonal(rng::AbstractRNG, d1::Integer, ds::Integer...; kwargs...)
313313
return reshape(orthogonal(rng, rows, cols; kwargs...), dims)
314314
end
315315

316-
orthogonal(dims::Integer...; kwargs...) = orthogonal(rng_from_array(), dims...; kwargs...)
317-
orthogonal(rng::AbstractRNG=rng_from_array(); init_kwargs...) = (dims::Integer...; kwargs...) -> orthogonal(rng, dims...; init_kwargs..., kwargs...)
316+
orthogonal(dims::Integer...; kwargs...) = orthogonal(_rng_from_array(), dims...; kwargs...)
317+
orthogonal(rng::AbstractRNG=_rng_from_array(); init_kwargs...) = (dims::Integer...; kwargs...) -> orthogonal(rng, dims...; init_kwargs..., kwargs...)
318318

319319
ChainRulesCore.@non_differentiable orthogonal(::Any...)
320320

@@ -361,8 +361,8 @@ function sparse_init(rng::AbstractRNG, dims::Integer...; sparsity, std = 0.01)
361361
return mapslices(shuffle, sparse_array, dims=1)
362362
end
363363

364-
sparse_init(dims::Integer...; kwargs...) = sparse_init(rng_from_array(), dims...; kwargs...)
365-
sparse_init(rng::AbstractRNG=rng_from_array(); init_kwargs...) = (dims...; kwargs...) -> sparse_init(rng, dims...; init_kwargs..., kwargs...)
364+
sparse_init(dims::Integer...; kwargs...) = sparse_init(_rng_from_array(), dims...; kwargs...)
365+
sparse_init(rng::AbstractRNG=_rng_from_array(); init_kwargs...) = (dims...; kwargs...) -> sparse_init(rng, dims...; init_kwargs..., kwargs...)
366366

367367
ChainRulesCore.@non_differentiable sparse_init(::Any...)
368368

@@ -452,7 +452,7 @@ end
452452

453453
# For consistency, it accepts an RNG, but ignores it:
454454
identity_init(::AbstractRNG, dims::Integer...; kwargs...) = identity_init(dims...; kwargs...)
455-
identity_init(rng::AbstractRNG=rng_from_array(); init_kwargs...) = (args...;kwargs...) -> identity_init(rng, args...; init_kwargs..., kwargs...)
455+
identity_init(rng::AbstractRNG=_rng_from_array(); init_kwargs...) = (args...;kwargs...) -> identity_init(rng, args...; init_kwargs..., kwargs...)
456456

457457
ChainRulesCore.@non_differentiable identity_init(::Any...)
458458

@@ -461,33 +461,40 @@ zeros32(dims::Integer...) = Base.zeros(Float32, dims...)
461461

462462
"""
463463
ones32(size...) = ones(Float32, size...)
464-
zeros32(size...) = zeros(Float32, size...)
465464
466-
Return an `Array{Float32}` of the given `size`.
465+
Return an `Array{Float32}` of the given `size` filled with 1s.
467466
"""
468467
ones32(dims...) = Base.ones(Float32, dims...)
469468

470-
@doc @doc(ones32)
469+
"""
470+
zeros32(size...) = zeros(Float32, size...)
471+
472+
Return an `Array{Float32}` of the given `size` filled with 0s.
473+
"""
471474
zeros32(dims...) = Base.zeros(Float32, dims...)
472475

473476
"""
474477
rand32([rng], size...)
475-
randn32([rng], size...)
476478
477-
Return an `Array{Float32}` of the given `size`, filled like `rand` or `randn`.
479+
Return an `Array{Float32}` of the given `size`, filled like `rand`.
478480
When the size is not provided, `rand32(rng::AbstractRNG)` returns a function.
479481
"""
480482
rand32(dims::Integer...) = Base.rand(Float32, dims...)
481483
rand32(rng::AbstractRNG, dims::Integer...) = Base.rand(rng, Float32, dims...)
482484
rand32(rng::AbstractRNG) = (dims...,) -> Base.rand(rng, Float32, dims...)
483485

484-
@doc @doc(rand32)
486+
"""
487+
randn32([rng], size...)
488+
489+
Return an `Array{Float32}` of the given `size`, filled like `randn`.
490+
When the size is not provided, `randn32(rng::AbstractRNG)` returns a function.
491+
"""
485492
randn32(dims::Integer...) = Base.randn(Float32, dims...)
486493
randn32(rng::AbstractRNG, dims::Integer...) = Base.randn(rng, Float32, dims...)
487494
randn32(rng::AbstractRNG) = (dims...,) -> Base.randn(rng, Float32, dims...)
488495

489496
"""
490-
create_bias(weights, bias, size...)
497+
_create_bias(weights, bias, size...)
491498
492499
Return a bias parameter for a layer, based on the value given
493500
to the constructor's keyword `bias=bias`.
@@ -497,10 +504,10 @@ to the constructor's keyword `bias=bias`.
497504
* `bias::AbstractArray` uses the array provided, provided it has the correct size.
498505
It does not at present correct the `eltype` to match that of `weights`.
499506
"""
500-
function create_bias(weights::AbstractArray, bias::Bool, dims::Integer...)
507+
function _create_bias(weights::AbstractArray, bias::Bool, dims::Integer...)
501508
bias ? fill!(similar(weights, dims...), 0) : false
502509
end
503-
function create_bias(weights::AbstractArray, bias::AbstractArray, dims::Integer...)
510+
function _create_bias(weights::AbstractArray, bias::AbstractArray, dims::Integer...)
504511
size(bias) == dims || throw(DimensionMismatch("expected bias of size $(dims), got size $(size(bias))"))
505512
bias
506513
end
@@ -518,6 +525,34 @@ Normally, the throttled function will run as much as it can, without ever
518525
going more than once per `wait` duration; but if you'd like to disable the
519526
execution on the leading edge, pass `leading=false`. To enable execution on
520527
the trailing edge, pass `trailing=true`.
528+
529+
# Examples
530+
```jldoctest
531+
julia> a = Flux.throttle(() -> println("Flux"), 2);
532+
533+
julia> a()
534+
Flux
535+
536+
julia> a()
537+
Flux
538+
539+
julia> for i = 1:4 # sleeps for 1 second -> the function can be called in alternate iterations
540+
a()
541+
sleep(1)
542+
end
543+
Flux
544+
Flux
545+
546+
julia> for i = 1:4 # sleeps for 2 second -> the function can be called in the next iteration
547+
a()
548+
sleep(2)
549+
end
550+
Flux
551+
Flux
552+
Flux
553+
Flux
554+
555+
```
521556
"""
522557
function throttle(f, timeout; leading=true, trailing=false)
523558
cooldown = true

0 commit comments

Comments
 (0)